deep-person-reid/train_img_model_xent.py

252 lines
9.1 KiB
Python
Raw Normal View History

2018-03-12 05:17:48 +08:00
from __future__ import absolute_import
import os
import sys
import time
import datetime
import argparse
import os.path as osp
import numpy as np
import torch
2018-03-12 05:27:48 +08:00
import torch.nn as nn
2018-03-12 05:17:48 +08:00
import torch.backends.cudnn as cudnn
from torch.utils.data import DataLoader
from torch.autograd import Variable
2018-03-12 06:52:33 +08:00
from torch.optim import lr_scheduler
2018-03-12 05:17:48 +08:00
import data_manager
from dataset_loader import ImageDataset
import transforms as T
import models
from losses import CrossEntropyLabelSmooth
2018-03-12 07:24:47 +08:00
from utils import AverageMeter, Logger, save_checkpoint
2018-03-12 05:17:48 +08:00
from eval_metrics import evaluate
parser = argparse.ArgumentParser(description='Train image model with cross entropy loss')
# Datasets
parser.add_argument('-d', '--dataset', type=str, default='market1501',
choices=data_manager.get_names())
parser.add_argument('-j', '--workers', default=4, type=int,
help="number of data loading workers (default: 4)")
parser.add_argument('--height', type=int, default=256,
help="height of an image (default: 256)")
parser.add_argument('--width', type=int, default=128,
help="width of an image (default: 128)")
# Optimization options
2018-03-12 06:52:33 +08:00
parser.add_argument('--max-epoch', default=60, type=int,
2018-03-12 05:17:48 +08:00
help="maximum epochs to run")
parser.add_argument('--start-epoch', default=0, type=int,
help="manual epoch number (useful on restarts)")
parser.add_argument('--train-batch', default=32, type=int,
help="train batch size")
2018-03-12 21:53:08 +08:00
parser.add_argument('--test-batch', default=32, type=int, help="test batch size")
2018-03-12 07:24:47 +08:00
parser.add_argument('--lr', '--learning-rate', default=0.0003, type=float,
2018-03-12 05:17:48 +08:00
help="initial learning rate")
2018-03-12 07:24:47 +08:00
parser.add_argument('--stepsize', default=20, type=int,
2018-03-12 06:52:33 +08:00
help="stepsize to decay learning rate (>0 means this is enabled)")
parser.add_argument('--gamma', default=0.1, type=float,
help="learning rate decay")
2018-03-12 07:24:47 +08:00
parser.add_argument('--weight-decay', default=5e-04, type=float,
2018-03-12 05:17:48 +08:00
help="weight decay (default: 5e-04)")
# Architecture
parser.add_argument('-a', '--arch', type=str, default='resnet50', choices=models.get_names())
# Miscs
2018-03-12 07:24:47 +08:00
parser.add_argument('--print-freq', type=int, default=10, help="print frequency")
2018-03-12 05:17:48 +08:00
parser.add_argument('--seed', type=int, default=1, help="manual seed")
parser.add_argument('--resume', type=str, default='', metavar='PATH')
parser.add_argument('--evaluate', action='store_true', help="evaluation only")
2018-03-12 07:24:47 +08:00
parser.add_argument('--eval-step', type=int, default=-1,
help="run evaluation for every N epochs (set to -1 to test after training)")
2018-03-12 05:17:48 +08:00
parser.add_argument('--save-dir', type=str, default='log')
parser.add_argument('--use-cpu', action='store_true', help="use cpu")
parser.add_argument('--gpu-devices', default='0', type=str, help='gpu device ids for CUDA_VISIBLE_DEVICES')
args = parser.parse_args()
def main():
torch.manual_seed(args.seed)
os.environ['CUDA_VISIBLE_DEVICES'] = args.gpu_devices
use_gpu = torch.cuda.is_available()
if args.use_cpu: use_gpu = False
2018-03-12 05:55:13 +08:00
if not args.evaluate:
sys.stdout = Logger(osp.join(args.save_dir, 'log_train.txt'))
else:
sys.stdout = Logger(osp.join(args.save_dir, 'log_test.txt'))
2018-03-12 05:42:57 +08:00
print("==========\nArgs:{}\n==========".format(args))
2018-03-12 05:35:49 +08:00
2018-03-12 05:17:48 +08:00
if use_gpu:
print("Currently using GPU")
cudnn.benchmark = True
torch.cuda.manual_seed_all(args.seed)
else:
print("Currently using CPU (GPU is highly recommended)")
print("Initializing dataset {}".format(args.dataset))
dataset = data_manager.init_dataset(name=args.dataset)
transform_train = T.Compose([
T.Random2DTranslation(args.height, args.width),
T.RandomHorizontalFlip(),
T.ToTensor(),
T.Normalize(mean=[0.485, 0.456, 0.406], std=[0.229, 0.224, 0.225]),
])
transform_test = T.Compose([
T.Resize((args.height, args.width)),
T.ToTensor(),
T.Normalize(mean=[0.485, 0.456, 0.406], std=[0.229, 0.224, 0.225]),
])
2018-03-12 21:18:25 +08:00
pin_memory = True if use_gpu else False
2018-03-12 05:17:48 +08:00
trainloader = DataLoader(
ImageDataset(dataset.train, transform=transform_train),
batch_size=args.train_batch, shuffle=True, num_workers=args.workers,
2018-03-12 21:18:25 +08:00
pin_memory=pin_memory, drop_last=True,
2018-03-12 05:17:48 +08:00
)
queryloader = DataLoader(
ImageDataset(dataset.query, transform=transform_test),
batch_size=args.test_batch, shuffle=False, num_workers=args.workers,
2018-03-12 21:18:25 +08:00
pin_memory=pin_memory, drop_last=False,
2018-03-12 05:17:48 +08:00
)
galleryloader = DataLoader(
ImageDataset(dataset.gallery, transform=transform_test),
batch_size=args.test_batch, shuffle=False, num_workers=args.workers,
2018-03-12 21:18:25 +08:00
pin_memory=pin_memory, drop_last=False,
2018-03-12 05:17:48 +08:00
)
print("Initializing model: {}".format(args.arch))
2018-03-12 21:53:08 +08:00
model = models.init_model(name=args.arch, num_classes=dataset.num_train_pids, loss={'xent'})
2018-03-12 05:17:48 +08:00
print("Model size: {:.5f}M".format(sum(p.numel() for p in model.parameters())/1000000.0))
criterion = CrossEntropyLabelSmooth(num_classes=dataset.num_train_pids, use_gpu=use_gpu)
optimizer = torch.optim.Adam(model.parameters(), lr=args.lr, weight_decay=args.weight_decay)
2018-03-12 06:52:33 +08:00
if args.stepsize > 0:
2018-03-12 06:54:02 +08:00
scheduler = lr_scheduler.StepLR(optimizer, step_size=args.stepsize, gamma=args.gamma)
2018-03-12 05:17:48 +08:00
start_epoch = args.start_epoch
if args.resume:
2018-03-12 06:29:53 +08:00
print("Loading checkpoint from '{}'".format(args.resume))
2018-03-12 05:17:48 +08:00
checkpoint = torch.load(args.resume)
model.load_state_dict(checkpoint['state_dict'])
start_epoch = checkpoint['epoch']
if use_gpu:
model = nn.DataParallel(model).cuda()
if args.evaluate:
print("Evaluate only")
2018-03-12 05:25:26 +08:00
test(model, queryloader, galleryloader, use_gpu)
2018-03-12 05:17:48 +08:00
return
start_time = time.time()
best_rank1 = -np.inf
for epoch in range(start_epoch, args.max_epoch):
print("==> Epoch {}/{}".format(epoch+1, args.max_epoch))
2018-03-12 06:52:33 +08:00
2018-03-12 05:17:48 +08:00
train(model, criterion, optimizer, trainloader, use_gpu)
2018-03-12 06:52:33 +08:00
if args.stepsize > 0: scheduler.step()
2018-03-12 07:24:47 +08:00
if args.eval_step > 0 and (epoch+1) % args.eval_step == 0 or (epoch+1) == args.max_epoch:
2018-03-12 05:17:48 +08:00
print("==> Test")
rank1 = test(model, queryloader, galleryloader, use_gpu)
is_best = rank1 > best_rank1
if is_best: best_rank1 = rank1
save_checkpoint({
'state_dict': model.state_dict(),
'rank1': rank1,
'epoch': epoch,
}, is_best, osp.join(args.save_dir, 'checkpoint_ep' + str(epoch+1) + '.pth.tar'))
2018-03-12 06:29:53 +08:00
elapsed = round(time.time() - start_time)
2018-03-12 05:17:48 +08:00
elapsed = str(datetime.timedelta(seconds=elapsed))
2018-03-12 06:29:53 +08:00
print("Finished. Total elapsed time (h:m:s): {}".format(elapsed))
2018-03-12 05:17:48 +08:00
def train(model, criterion, optimizer, trainloader, use_gpu):
model.train()
losses = AverageMeter()
for batch_idx, (imgs, pids, _) in enumerate(trainloader):
if use_gpu:
imgs, pids = imgs.cuda(), pids.cuda()
imgs, pids = Variable(imgs), Variable(pids)
outputs = model(imgs)
loss = criterion(outputs, pids)
optimizer.zero_grad()
loss.backward()
optimizer.step()
losses.update(loss.data[0], pids.size(0))
if (batch_idx+1) % args.print_freq == 0:
print("Batch {}/{}\t Loss {:.6f} ({:.6f})".format(batch_idx+1, len(trainloader), losses.val, losses.avg))
2018-03-12 06:10:43 +08:00
def test(model, queryloader, galleryloader, use_gpu, ranks=[1, 5, 10, 20]):
2018-03-12 05:17:48 +08:00
model.eval()
2018-03-12 06:29:53 +08:00
qf, q_pids, q_camids = [], [], []
for batch_idx, (imgs, pids, camids) in enumerate(queryloader):
2018-03-12 05:17:48 +08:00
if use_gpu:
imgs = imgs.cuda()
2018-03-13 18:23:44 +08:00
imgs = Variable(imgs, volatile=True)
2018-03-12 05:17:48 +08:00
features = model(imgs)
features = features.data.cpu()
qf.append(features)
2018-03-12 06:29:53 +08:00
q_pids.extend(pids)
q_camids.extend(camids)
2018-03-12 05:17:48 +08:00
qf = torch.cat(qf, 0)
2018-03-12 06:29:53 +08:00
q_pids = np.asarray(q_pids)
q_camids = np.asarray(q_camids)
2018-03-12 05:17:48 +08:00
2018-03-12 05:29:35 +08:00
print("Extracted features for query set, obtained {}-by-{} matrix".format(qf.size(0), qf.size(1)))
2018-03-12 05:17:48 +08:00
2018-03-12 06:29:53 +08:00
gf, g_pids, g_camids = [], [], []
for batch_idx, (imgs, pids, camids) in enumerate(galleryloader):
2018-03-12 05:17:48 +08:00
if use_gpu:
imgs = imgs.cuda()
2018-03-13 18:23:44 +08:00
imgs = Variable(imgs, volatile=True)
2018-03-12 05:17:48 +08:00
features = model(imgs)
features = features.data.cpu()
gf.append(features)
2018-03-12 06:29:53 +08:00
g_pids.extend(pids)
g_camids.extend(camids)
2018-03-12 05:17:48 +08:00
gf = torch.cat(gf, 0)
2018-03-12 06:29:53 +08:00
g_pids = np.asarray(g_pids)
g_camids = np.asarray(g_camids)
2018-03-12 05:17:48 +08:00
2018-03-12 05:29:35 +08:00
print("Extracted features for gallery set, obtained {}-by-{} matrix".format(gf.size(0), gf.size(1)))
2018-03-12 05:17:48 +08:00
print("Computing distance matrix")
m, n = qf.size(0), gf.size(0)
distmat = torch.pow(qf, 2).sum(dim=1, keepdim=True).expand(m, n) + \
torch.pow(gf, 2).sum(dim=1, keepdim=True).expand(n, m).t()
distmat.addmm_(1, -2, qf, gf.t())
2018-03-12 06:29:53 +08:00
distmat = distmat.numpy()
2018-03-12 05:17:48 +08:00
print("Computing CMC and mAP")
cmc, mAP = evaluate(distmat, q_pids, g_pids, q_camids, g_camids)
2018-03-12 06:36:46 +08:00
print("Results ----------")
print("mAP: {:.1%}".format(mAP))
print("CMC curve")
2018-03-12 06:10:43 +08:00
for r in ranks:
2018-03-12 06:20:10 +08:00
print("Rank-{:<3}: {:.1%}".format(r, cmc[r-1]))
2018-03-12 06:36:46 +08:00
print("------------------")
2018-03-12 05:17:48 +08:00
return cmc[0]
if __name__ == '__main__':
main()