engine.py 22 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558
  1. import torch
  2. import torch.distributed as dist
  3. import os
  4. import random
  5. # ----------------- Extra Components -----------------
  6. from utils import distributed_utils
  7. from utils.misc import MetricLogger, SmoothedValue
  8. from utils.vis_tools import vis_data
  9. # ----------------- Optimizer & LrScheduler Components -----------------
  10. from utils.solver.optimizer import build_yolo_optimizer, build_rtdetr_optimizer
  11. from utils.solver.lr_scheduler import LinearWarmUpLrScheduler, build_lr_scheduler
  12. class YoloTrainer(object):
  13. def __init__(self,
  14. # Basic parameters
  15. args,
  16. cfg,
  17. device,
  18. # Model parameters
  19. model,
  20. model_ema,
  21. criterion,
  22. # Data parameters
  23. train_transform,
  24. val_transform,
  25. dataset,
  26. train_loader,
  27. evaluator,
  28. ):
  29. # ------------------- basic parameters -------------------
  30. self.args = args
  31. self.cfg = cfg
  32. self.epoch = 0
  33. self.best_map = -1.
  34. self.device = device
  35. self.criterion = criterion
  36. self.heavy_eval = False
  37. self.model_ema = model_ema
  38. # weak augmentatino stage
  39. self.second_stage = False
  40. self.second_stage_epoch = cfg.no_aug_epoch
  41. # path to save model
  42. self.path_to_save = os.path.join(args.save_folder, args.dataset, args.model)
  43. os.makedirs(self.path_to_save, exist_ok=True)
  44. # ---------------------------- Transform ----------------------------
  45. self.train_transform = train_transform
  46. self.val_transform = val_transform
  47. # ---------------------------- Dataset & Dataloader ----------------------------
  48. self.dataset = dataset
  49. self.train_loader = train_loader
  50. # ---------------------------- Evaluator ----------------------------
  51. self.evaluator = evaluator
  52. # ---------------------------- Build Grad. Scaler ----------------------------
  53. self.scaler = torch.cuda.amp.GradScaler(enabled=args.fp16)
  54. # ---------------------------- Build Optimizer ----------------------------
  55. self.grad_accumulate = max(64 // args.batch_size, 1)
  56. cfg.base_lr = cfg.per_image_lr * args.batch_size * self.grad_accumulate
  57. cfg.min_lr = cfg.base_lr * cfg.min_lr_ratio
  58. self.optimizer, self.start_epoch = build_yolo_optimizer(cfg, model, args.resume)
  59. # ---------------------------- Build LR Scheduler ----------------------------
  60. warmup_iters = cfg.warmup_epoch * len(self.train_loader)
  61. self.lr_scheduler_warmup = LinearWarmUpLrScheduler(warmup_iters, cfg.base_lr, cfg.warmup_bias_lr)
  62. self.lr_scheduler = build_lr_scheduler(cfg, self.optimizer, args.resume)
  63. self.best_map = cfg.best_map / 100.0
  64. print("Best mAP metric: {}".format(self.best_map))
  65. def train(self, model):
  66. for epoch in range(self.start_epoch, self.cfg.max_epoch):
  67. if self.args.distributed:
  68. self.train_loader.batch_sampler.sampler.set_epoch(epoch)
  69. # check second stage
  70. if epoch >= (self.cfg.max_epoch - self.second_stage_epoch - 1) and not self.second_stage:
  71. self.check_second_stage()
  72. # save model of the last mosaic epoch
  73. weight_name = '{}_last_mosaic_epoch.pth'.format(self.args.model)
  74. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  75. print('Saving state of the last Mosaic epoch-{}.'.format(self.epoch))
  76. torch.save({'model': model.state_dict(),
  77. 'mAP': round(self.evaluator.map*100, 1),
  78. 'optimizer': self.optimizer.state_dict(),
  79. 'epoch': self.epoch,
  80. 'args': self.args},
  81. checkpoint_path)
  82. # train one epoch
  83. self.epoch = epoch
  84. self.train_one_epoch(model)
  85. # LR Schedule
  86. if (epoch + 1) > self.cfg.warmup_epoch:
  87. self.lr_scheduler.step()
  88. # eval one epoch
  89. if self.heavy_eval:
  90. model_eval = model.module if self.args.distributed else model
  91. self.eval(model_eval)
  92. else:
  93. model_eval = model.module if self.args.distributed else model
  94. if (epoch % self.cfg.eval_epoch) == 0 or (epoch == self.cfg.max_epoch - 1):
  95. self.eval(model_eval)
  96. if self.args.debug:
  97. print("For debug mode, we only train 1 epoch")
  98. break
  99. def eval(self, model):
  100. # set eval mode
  101. model.eval()
  102. model_eval = model if self.model_ema is None else self.model_ema.ema
  103. cur_map = -1.
  104. to_save = False
  105. if distributed_utils.is_main_process():
  106. if self.evaluator is None:
  107. print('No evaluator ... save model and go on training.')
  108. to_save = True
  109. weight_name = '{}_no_eval.pth'.format(self.args.model)
  110. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  111. else:
  112. print('Eval ...')
  113. # Evaluate
  114. with torch.no_grad():
  115. self.evaluator.evaluate(model_eval)
  116. cur_map = self.evaluator.map
  117. if cur_map > self.best_map:
  118. # update best-map
  119. self.best_map = cur_map
  120. to_save = True
  121. # Save model
  122. if to_save:
  123. print('Saving state, epoch:', self.epoch)
  124. weight_name = '{}_best.pth'.format(self.args.model)
  125. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  126. state_dicts = {
  127. 'model': model_eval.state_dict(),
  128. 'mAP': round(cur_map*100, 3),
  129. 'optimizer': self.optimizer.state_dict(),
  130. 'lr_scheduler': self.lr_scheduler.state_dict(),
  131. 'epoch': self.epoch,
  132. 'args': self.args,
  133. }
  134. if self.model_ema is not None:
  135. state_dicts["ema_updates"] = self.model_ema.updates
  136. torch.save(state_dicts, checkpoint_path)
  137. if self.args.distributed:
  138. # wait for all processes to synchronize
  139. dist.barrier()
  140. # set train mode.
  141. model.train()
  142. def train_one_epoch(self, model):
  143. metric_logger = MetricLogger(delimiter=" ")
  144. metric_logger.add_meter('lr', SmoothedValue(window_size=1, fmt='{value:.6f}'))
  145. metric_logger.add_meter('size', SmoothedValue(window_size=1, fmt='{value:d}'))
  146. metric_logger.add_meter('gnorm', SmoothedValue(window_size=1, fmt='{value:.1f}'))
  147. header = 'Epoch: [{} / {}]'.format(self.epoch, self.cfg.max_epoch)
  148. epoch_size = len(self.train_loader)
  149. print_freq = 10
  150. gnorm = 0.0
  151. # basic parameters
  152. epoch_size = len(self.train_loader)
  153. img_size = self.cfg.train_img_size
  154. nw = epoch_size * self.cfg.warmup_epoch
  155. # Train one epoch
  156. for iter_i, (images, targets) in enumerate(metric_logger.log_every(self.train_loader, print_freq, header)):
  157. ni = iter_i + self.epoch * epoch_size
  158. # Warmup
  159. if nw > 0 and ni < nw:
  160. self.lr_scheduler_warmup(ni, self.optimizer)
  161. elif ni == nw:
  162. print("Warmup stage is over.")
  163. self.lr_scheduler_warmup.set_lr(self.optimizer, self.cfg.base_lr)
  164. # To device
  165. images = images.to(self.device, non_blocking=True).float()
  166. # Multi scale
  167. images, targets, img_size = self.rescale_image_targets(
  168. images, targets, self.cfg.max_stride, self.cfg.multi_scale)
  169. # Visualize train targets
  170. if self.args.vis_tgt:
  171. vis_data(images,
  172. targets,
  173. self.cfg.num_classes,
  174. self.cfg.normalize_coords,
  175. self.train_transform.color_format,
  176. self.cfg.pixel_mean,
  177. self.cfg.pixel_std,
  178. self.cfg.box_format)
  179. # Inference
  180. with torch.cuda.amp.autocast(enabled=self.args.fp16):
  181. outputs = model(images)
  182. # Compute loss
  183. loss_dict = self.criterion(outputs=outputs, targets=targets)
  184. losses = loss_dict['losses']
  185. losses /= self.grad_accumulate
  186. loss_dict_reduced = distributed_utils.reduce_dict(loss_dict)
  187. # Backward
  188. self.scaler.scale(losses).backward()
  189. # Optimize
  190. if (iter_i + 1) % self.grad_accumulate == 0:
  191. if self.cfg.clip_max_norm > 0:
  192. self.scaler.unscale_(self.optimizer)
  193. gnorm = torch.nn.utils.clip_grad_norm_(model.parameters(), max_norm=self.cfg.clip_max_norm)
  194. self.scaler.step(self.optimizer)
  195. self.scaler.update()
  196. self.optimizer.zero_grad()
  197. # ModelEMA
  198. if self.model_ema is not None:
  199. self.model_ema.update(model)
  200. # Update log
  201. metric_logger.update(**loss_dict_reduced)
  202. metric_logger.update(lr=self.optimizer.param_groups[2]["lr"])
  203. metric_logger.update(size=img_size)
  204. metric_logger.update(gnorm=gnorm)
  205. if self.args.debug:
  206. print("For debug mode, we only train 1 iteration")
  207. break
  208. # Gather the stats from all processes
  209. metric_logger.synchronize_between_processes()
  210. print("Averaged stats:", metric_logger)
  211. def rescale_image_targets(self, images, targets, max_stride, multi_scale_range=[0.5, 1.5]):
  212. """
  213. Deployed for Multi scale trick.
  214. """
  215. # During training phase, the shape of input image is square.
  216. old_img_size = images.shape[-1]
  217. min_img_size = old_img_size * multi_scale_range[0]
  218. max_img_size = old_img_size * multi_scale_range[1]
  219. # Choose a new image size
  220. new_img_size = random.randrange(min_img_size, max_img_size + max_stride, max_stride)
  221. # Resize
  222. if new_img_size != old_img_size:
  223. # interpolate
  224. images = torch.nn.functional.interpolate(
  225. input=images,
  226. size=new_img_size,
  227. mode='bilinear',
  228. align_corners=False)
  229. # rescale targets
  230. if not self.cfg.normalize_coords:
  231. for tgt in targets:
  232. boxes = tgt["boxes"].clone()
  233. labels = tgt["labels"].clone()
  234. boxes = torch.clamp(boxes, 0, old_img_size)
  235. # rescale box
  236. boxes[:, [0, 2]] = boxes[:, [0, 2]] / old_img_size * new_img_size
  237. boxes[:, [1, 3]] = boxes[:, [1, 3]] / old_img_size * new_img_size
  238. # refine tgt
  239. tgt_boxes_wh = boxes[..., 2:] - boxes[..., :2]
  240. min_tgt_size = torch.min(tgt_boxes_wh, dim=-1)[0]
  241. keep = (min_tgt_size >= 8)
  242. tgt["boxes"] = boxes[keep]
  243. tgt["labels"] = labels[keep]
  244. return images, targets, new_img_size
  245. def check_second_stage(self):
  246. # set second stage
  247. print('============== Second stage of Training ==============')
  248. self.second_stage = True
  249. self.heavy_eval = True
  250. # close mosaic augmentation
  251. if self.train_loader.dataset.mosaic_prob > 0.:
  252. print(' - Close < Mosaic Augmentation > ...')
  253. self.train_loader.dataset.mosaic_prob = 0.
  254. # close mixup augmentation
  255. if self.train_loader.dataset.mixup_prob > 0.:
  256. print(' - Close < Mixup Augmentation > ...')
  257. self.train_loader.dataset.mixup_prob = 0.
  258. # close copy-paste augmentation
  259. if self.train_loader.dataset.copy_paste > 0.:
  260. print(' - Close < Copy-paste Augmentation > ...')
  261. self.train_loader.dataset.copy_paste = 0.
  262. class RTDetrTrainer(object):
  263. def __init__(self,
  264. # Basic parameters
  265. args,
  266. cfg,
  267. device,
  268. # Model parameters
  269. model,
  270. model_ema,
  271. criterion,
  272. # Data parameters
  273. train_transform,
  274. val_transform,
  275. dataset,
  276. train_loader,
  277. evaluator,
  278. ):
  279. # ------------------- basic parameters -------------------
  280. self.args = args
  281. self.cfg = cfg
  282. self.epoch = 0
  283. self.best_map = -1.
  284. self.device = device
  285. self.criterion = criterion
  286. self.heavy_eval = False
  287. self.model_ema = model_ema
  288. # path to save model
  289. self.path_to_save = os.path.join(args.save_folder, args.dataset, args.model)
  290. os.makedirs(self.path_to_save, exist_ok=True)
  291. # ---------------------------- Transform ----------------------------
  292. self.train_transform = train_transform
  293. self.val_transform = val_transform
  294. # ---------------------------- Dataset & Dataloader ----------------------------
  295. self.dataset = dataset
  296. self.train_loader = train_loader
  297. # ---------------------------- Evaluator ----------------------------
  298. self.evaluator = evaluator
  299. # ---------------------------- Build Grad. Scaler ----------------------------
  300. self.scaler = torch.cuda.amp.GradScaler(enabled=args.fp16)
  301. # ---------------------------- Build Optimizer ----------------------------
  302. self.grad_accumulate = max(16 // args.batch_size, 1)
  303. cfg.base_lr = cfg.per_image_lr * args.batch_size * self.grad_accumulate
  304. cfg.min_lr = cfg.base_lr * cfg.min_lr_ratio
  305. self.optimizer, self.start_epoch = build_rtdetr_optimizer(cfg, model, args.resume)
  306. # ---------------------------- Build LR Scheduler ----------------------------
  307. self.wp_lr_scheduler = LinearWarmUpLrScheduler(cfg.warmup_iters, cfg.base_lr)
  308. self.lr_scheduler = build_lr_scheduler(cfg, self.optimizer, args.resume)
  309. def train(self, model):
  310. for epoch in range(self.start_epoch, self.cfg.max_epoch):
  311. if self.args.distributed:
  312. self.train_loader.batch_sampler.sampler.set_epoch(epoch)
  313. # train one epoch
  314. self.epoch = epoch
  315. self.train_one_epoch(model)
  316. # LR Scheduler
  317. self.lr_scheduler.step()
  318. # eval one epoch
  319. if self.heavy_eval:
  320. model_eval = model.module if self.args.distributed else model
  321. self.eval(model_eval)
  322. else:
  323. model_eval = model.module if self.args.distributed else model
  324. if (epoch % self.cfg.eval_epoch) == 0 or (epoch == self.cfg.max_epoch - 1):
  325. self.eval(model_eval)
  326. if self.args.debug:
  327. print("For debug mode, we only train 1 epoch")
  328. break
  329. def eval(self, model):
  330. # set eval mode
  331. model.eval()
  332. model_eval = model if self.model_ema is None else self.model_ema.ema
  333. cur_map = -1.
  334. to_save = False
  335. if distributed_utils.is_main_process():
  336. if self.evaluator is None:
  337. print('No evaluator ... save model and go on training.')
  338. to_save = True
  339. weight_name = '{}_no_eval.pth'.format(self.args.model)
  340. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  341. else:
  342. print('Eval ...')
  343. # Evaluate
  344. with torch.no_grad():
  345. self.evaluator.evaluate(model_eval)
  346. cur_map = self.evaluator.map
  347. if cur_map > self.best_map:
  348. # update best-map
  349. self.best_map = cur_map
  350. to_save = True
  351. # Save model
  352. if to_save:
  353. print('Saving state, epoch:', self.epoch)
  354. weight_name = '{}_best.pth'.format(self.args.model)
  355. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  356. state_dicts = {
  357. 'model': model_eval.state_dict(),
  358. 'mAP': round(cur_map*100, 1),
  359. 'optimizer': self.optimizer.state_dict(),
  360. 'lr_scheduler': self.lr_scheduler.state_dict(),
  361. 'epoch': self.epoch,
  362. 'args': self.args,
  363. }
  364. if self.model_ema is not None:
  365. state_dicts["ema_updates"] = self.model_ema.updates
  366. torch.save(state_dicts, checkpoint_path)
  367. if self.args.distributed:
  368. # wait for all processes to synchronize
  369. dist.barrier()
  370. # set train mode.
  371. model.train()
  372. def train_one_epoch(self, model):
  373. metric_logger = MetricLogger(delimiter=" ")
  374. metric_logger.add_meter('lr', SmoothedValue(window_size=1, fmt='{value:.6f}'))
  375. metric_logger.add_meter('size', SmoothedValue(window_size=1, fmt='{value:d}'))
  376. metric_logger.add_meter('gnorm', SmoothedValue(window_size=1, fmt='{value:.1f}'))
  377. header = 'Epoch: [{} / {}]'.format(self.epoch, self.cfg.max_epoch)
  378. epoch_size = len(self.train_loader)
  379. print_freq = 10
  380. gnorm = 0.0
  381. # basic parameters
  382. epoch_size = len(self.train_loader)
  383. img_size = self.cfg.train_img_size
  384. nw = self.cfg.warmup_iters
  385. lr_warmup_stage = True
  386. # Train one epoch
  387. for iter_i, (images, targets) in enumerate(metric_logger.log_every(self.train_loader, print_freq, header)):
  388. ni = iter_i + self.epoch * epoch_size
  389. # WarmUp
  390. if lr_warmup_stage:
  391. if ni % self.grad_accumulate == 0:
  392. ni = ni // self.grad_accumulate
  393. if ni < nw:
  394. self.wp_lr_scheduler(ni, self.optimizer)
  395. elif ni == nw and lr_warmup_stage:
  396. print('Warmup stage is over.')
  397. lr_warmup_stage = False
  398. self.wp_lr_scheduler.set_lr(self.optimizer, self.cfg.base_lr)
  399. # To device
  400. images = images.to(self.device, non_blocking=True).float()
  401. for tgt in targets:
  402. tgt['boxes'] = tgt['boxes'].to(self.device)
  403. tgt['labels'] = tgt['labels'].to(self.device)
  404. # Multi scale
  405. images, targets, img_size = self.rescale_image_targets(
  406. images, targets, self.cfg.max_stride, self.cfg.multi_scale)
  407. # Visualize train targets
  408. if self.args.vis_tgt:
  409. vis_data(images,
  410. targets,
  411. self.cfg.num_classes,
  412. self.cfg.normalize_coords,
  413. self.train_transform.color_format,
  414. self.cfg.pixel_mean,
  415. self.cfg.pixel_std,
  416. self.cfg.box_format)
  417. # Inference
  418. with torch.cuda.amp.autocast(enabled=self.args.fp16):
  419. outputs = model(images, targets)
  420. # Compute loss
  421. loss_dict = self.criterion(outputs, targets)
  422. losses = sum(loss_dict.values())
  423. losses /= self.grad_accumulate
  424. loss_dict_reduced = distributed_utils.reduce_dict(loss_dict)
  425. # Backward
  426. self.scaler.scale(losses).backward()
  427. # Optimize
  428. if (iter_i + 1) % self.grad_accumulate == 0:
  429. if self.cfg.clip_max_norm > 0:
  430. self.scaler.unscale_(self.optimizer)
  431. gnorm = torch.nn.utils.clip_grad_norm_(model.parameters(), max_norm=self.cfg.clip_max_norm)
  432. self.scaler.step(self.optimizer)
  433. self.scaler.update()
  434. self.optimizer.zero_grad()
  435. # ModelEMA
  436. if self.model_ema is not None:
  437. self.model_ema.update(model)
  438. # Update log
  439. metric_logger.update(**loss_dict_reduced)
  440. metric_logger.update(lr=self.optimizer.param_groups[2]["lr"])
  441. metric_logger.update(size=img_size)
  442. metric_logger.update(gnorm=gnorm)
  443. if self.args.debug:
  444. print("For debug mode, we only train 1 iteration")
  445. break
  446. def rescale_image_targets(self, images, targets, max_stride, multi_scale_range=[0.5, 1.5]):
  447. """
  448. Deployed for Multi scale trick.
  449. """
  450. # During training phase, the shape of input image is square.
  451. old_img_size = images.shape[-1]
  452. min_img_size = old_img_size * multi_scale_range[0]
  453. max_img_size = old_img_size * multi_scale_range[1]
  454. # Choose a new image size
  455. new_img_size = random.randrange(min_img_size, max_img_size + max_stride, max_stride)
  456. # Resize
  457. if new_img_size != old_img_size:
  458. # interpolate
  459. images = torch.nn.functional.interpolate(
  460. input=images,
  461. size=new_img_size,
  462. mode='bilinear',
  463. align_corners=False)
  464. return images, targets, new_img_size
  465. # Build Trainer
  466. def build_trainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator):
  467. # ----------------------- Det trainers -----------------------
  468. if cfg.trainer == 'yolo':
  469. return YoloTrainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator)
  470. elif cfg.trainer == 'rtdetr':
  471. return RTDetrTrainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator)
  472. else:
  473. raise NotImplementedError(cfg.trainer)