qdtrack.py 7.7 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186
  1. # Copyright (c) OpenMMLab. All rights reserved.
  2. from typing import Optional, Union
  3. import torch
  4. from torch import Tensor
  5. from mmdet.registry import MODELS
  6. from mmdet.structures import TrackSampleList
  7. from mmdet.utils import OptConfigType, OptMultiConfig
  8. from .base import BaseMOTModel
  9. @MODELS.register_module()
  10. class QDTrack(BaseMOTModel):
  11. """Quasi-Dense Similarity Learning for Multiple Object Tracking.
  12. This multi object tracker is the implementation of `QDTrack
  13. <https://arxiv.org/abs/2006.06664>`_.
  14. Args:
  15. detector (dict): Configuration of detector. Defaults to None.
  16. track_head (dict): Configuration of track head. Defaults to None.
  17. tracker (dict): Configuration of tracker. Defaults to None.
  18. freeze_detector (bool): If True, freeze the detector weights.
  19. Defaults to False.
  20. data_preprocessor (dict or ConfigDict, optional): The pre-process
  21. config of :class:`TrackDataPreprocessor`. it usually includes,
  22. ``pad_size_divisor``, ``pad_value``, ``mean`` and ``std``.
  23. init_cfg (dict or list[dict]): Configuration of initialization.
  24. Defaults to None.
  25. """
  26. def __init__(self,
  27. detector: Optional[dict] = None,
  28. track_head: Optional[dict] = None,
  29. tracker: Optional[dict] = None,
  30. freeze_detector: bool = False,
  31. data_preprocessor: OptConfigType = None,
  32. init_cfg: OptMultiConfig = None):
  33. super().__init__(data_preprocessor, init_cfg)
  34. if detector is not None:
  35. self.detector = MODELS.build(detector)
  36. if track_head is not None:
  37. self.track_head = MODELS.build(track_head)
  38. if tracker is not None:
  39. self.tracker = MODELS.build(tracker)
  40. self.freeze_detector = freeze_detector
  41. if self.freeze_detector:
  42. self.freeze_module('detector')
  43. def predict(self,
  44. inputs: Tensor,
  45. data_samples: TrackSampleList,
  46. rescale: bool = True,
  47. **kwargs) -> TrackSampleList:
  48. """Predict results from a video and data samples with post- processing.
  49. Args:
  50. inputs (Tensor): of shape (N, T, C, H, W) encoding
  51. input images. The N denotes batch size.
  52. The T denotes the number of frames in a video.
  53. data_samples (list[:obj:`TrackDataSample`]): The batch
  54. data samples. It usually includes information such
  55. as `video_data_samples`.
  56. rescale (bool, Optional): If False, then returned bboxes and masks
  57. will fit the scale of img, otherwise, returned bboxes and masks
  58. will fit the scale of original image shape. Defaults to True.
  59. Returns:
  60. TrackSampleList: Tracking results of the inputs.
  61. """
  62. assert inputs.dim() == 5, 'The img must be 5D Tensor (N, T, C, H, W).'
  63. assert inputs.size(0) == 1, \
  64. 'QDTrack inference only support 1 batch size per gpu for now.'
  65. assert len(data_samples) == 1, \
  66. 'QDTrack only support 1 batch size per gpu for now.'
  67. track_data_sample = data_samples[0]
  68. video_len = len(track_data_sample)
  69. if track_data_sample[0].frame_id == 0:
  70. self.tracker.reset()
  71. for frame_id in range(video_len):
  72. img_data_sample = track_data_sample[frame_id]
  73. single_img = inputs[:, frame_id].contiguous()
  74. x = self.detector.extract_feat(single_img)
  75. rpn_results_list = self.detector.rpn_head.predict(
  76. x, [img_data_sample])
  77. # det_results List[InstanceData]
  78. det_results = self.detector.roi_head.predict(
  79. x, rpn_results_list, [img_data_sample], rescale=rescale)
  80. assert len(det_results) == 1, 'Batch inference is not supported.'
  81. img_data_sample.pred_instances = det_results[0]
  82. frame_pred_track_instances = self.tracker.track(
  83. model=self,
  84. img=single_img,
  85. feats=x,
  86. data_sample=img_data_sample,
  87. **kwargs)
  88. img_data_sample.pred_track_instances = frame_pred_track_instances
  89. return [track_data_sample]
  90. def loss(self, inputs: Tensor, data_samples: TrackSampleList,
  91. **kwargs) -> Union[dict, tuple]:
  92. """Calculate losses from a batch of inputs and data samples.
  93. Args:
  94. inputs (Dict[str, Tensor]): of shape (N, T, C, H, W) encoding
  95. input images. Typically these should be mean centered and std
  96. scaled. The N denotes batch size. The T denotes the number of
  97. frames.
  98. data_samples (list[:obj:`TrackDataSample`]): The batch
  99. data samples. It usually includes information such
  100. as `video_data_samples`.
  101. Returns:
  102. dict: A dictionary of loss components.
  103. """
  104. # modify the inputs shape to fit mmdet
  105. assert inputs.dim() == 5, 'The img must be 5D Tensor (N, T, C, H, W).'
  106. assert inputs.size(1) == 2, \
  107. 'QDTrack can only have 1 key frame and 1 reference frame.'
  108. # split the data_samples into two aspects: key frames and reference
  109. # frames
  110. ref_data_samples, key_data_samples = [], []
  111. key_frame_inds, ref_frame_inds = [], []
  112. # set cat_id of gt_labels to 0 in RPN
  113. for track_data_sample in data_samples:
  114. key_frame_inds.append(track_data_sample.key_frames_inds[0])
  115. ref_frame_inds.append(track_data_sample.ref_frames_inds[0])
  116. key_data_sample = track_data_sample.get_key_frames()[0]
  117. key_data_sample.gt_instances.labels = \
  118. torch.zeros_like(key_data_sample.gt_instances.labels)
  119. key_data_samples.append(key_data_sample)
  120. ref_data_sample = track_data_sample.get_ref_frames()[0]
  121. ref_data_samples.append(ref_data_sample)
  122. key_frame_inds = torch.tensor(key_frame_inds, dtype=torch.int64)
  123. ref_frame_inds = torch.tensor(ref_frame_inds, dtype=torch.int64)
  124. batch_inds = torch.arange(len(inputs))
  125. key_imgs = inputs[batch_inds, key_frame_inds].contiguous()
  126. ref_imgs = inputs[batch_inds, ref_frame_inds].contiguous()
  127. x = self.detector.extract_feat(key_imgs)
  128. ref_x = self.detector.extract_feat(ref_imgs)
  129. losses = dict()
  130. # RPN head forward and loss
  131. assert self.detector.with_rpn, \
  132. 'QDTrack only support detector with RPN.'
  133. proposal_cfg = self.detector.train_cfg.get('rpn_proposal',
  134. self.detector.test_cfg.rpn)
  135. rpn_losses, rpn_results_list = self.detector.rpn_head. \
  136. loss_and_predict(x,
  137. key_data_samples,
  138. proposal_cfg=proposal_cfg,
  139. **kwargs)
  140. ref_rpn_results_list = self.detector.rpn_head.predict(
  141. ref_x, ref_data_samples, **kwargs)
  142. # avoid get same name with roi_head loss
  143. keys = rpn_losses.keys()
  144. for key in keys:
  145. if 'loss' in key and 'rpn' not in key:
  146. rpn_losses[f'rpn_{key}'] = rpn_losses.pop(key)
  147. losses.update(rpn_losses)
  148. # roi_head loss
  149. losses_detect = self.detector.roi_head.loss(x, rpn_results_list,
  150. key_data_samples, **kwargs)
  151. losses.update(losses_detect)
  152. # tracking head loss
  153. losses_track = self.track_head.loss(x, ref_x, rpn_results_list,
  154. ref_rpn_results_list, data_samples,
  155. **kwargs)
  156. losses.update(losses_track)
  157. return losses