base_detr.py 14 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332
  1. # Copyright (c) OpenMMLab. All rights reserved.
  2. from abc import ABCMeta, abstractmethod
  3. from typing import Dict, List, Tuple, Union
  4. from torch import Tensor
  5. from mmdet.registry import MODELS
  6. from mmdet.structures import OptSampleList, SampleList
  7. from mmdet.utils import ConfigType, OptConfigType, OptMultiConfig
  8. from .base import BaseDetector
  9. @MODELS.register_module()
  10. class DetectionTransformer(BaseDetector, metaclass=ABCMeta):
  11. r"""Base class for Detection Transformer.
  12. In Detection Transformer, an encoder is used to process output features of
  13. neck, then several queries interact with the encoder features using a
  14. decoder and do the regression and classification with the bounding box
  15. head.
  16. Args:
  17. backbone (:obj:`ConfigDict` or dict): Config of the backbone.
  18. neck (:obj:`ConfigDict` or dict, optional): Config of the neck.
  19. Defaults to None.
  20. encoder (:obj:`ConfigDict` or dict, optional): Config of the
  21. Transformer encoder. Defaults to None.
  22. decoder (:obj:`ConfigDict` or dict, optional): Config of the
  23. Transformer decoder. Defaults to None.
  24. bbox_head (:obj:`ConfigDict` or dict, optional): Config for the
  25. bounding box head module. Defaults to None.
  26. positional_encoding (:obj:`ConfigDict` or dict, optional): Config
  27. of the positional encoding module. Defaults to None.
  28. num_queries (int, optional): Number of decoder query in Transformer.
  29. Defaults to 100.
  30. train_cfg (:obj:`ConfigDict` or dict, optional): Training config of
  31. the bounding box head module. Defaults to None.
  32. test_cfg (:obj:`ConfigDict` or dict, optional): Testing config of
  33. the bounding box head module. Defaults to None.
  34. data_preprocessor (dict or ConfigDict, optional): The pre-process
  35. config of :class:`BaseDataPreprocessor`. it usually includes,
  36. ``pad_size_divisor``, ``pad_value``, ``mean`` and ``std``.
  37. Defaults to None.
  38. init_cfg (:obj:`ConfigDict` or dict, optional): the config to control
  39. the initialization. Defaults to None.
  40. """
  41. def __init__(self,
  42. backbone: ConfigType,
  43. neck: OptConfigType = None,
  44. encoder: OptConfigType = None,
  45. decoder: OptConfigType = None,
  46. bbox_head: OptConfigType = None,
  47. positional_encoding: OptConfigType = None,
  48. num_queries: int = 100,
  49. train_cfg: OptConfigType = None,
  50. test_cfg: OptConfigType = None,
  51. data_preprocessor: OptConfigType = None,
  52. init_cfg: OptMultiConfig = None) -> None:
  53. super().__init__(
  54. data_preprocessor=data_preprocessor, init_cfg=init_cfg)
  55. # process args
  56. bbox_head.update(train_cfg=train_cfg)
  57. bbox_head.update(test_cfg=test_cfg)
  58. self.train_cfg = train_cfg
  59. self.test_cfg = test_cfg
  60. self.encoder = encoder
  61. self.decoder = decoder
  62. self.positional_encoding = positional_encoding
  63. self.num_queries = num_queries
  64. # init model layers
  65. self.backbone = MODELS.build(backbone)
  66. if neck is not None:
  67. self.neck = MODELS.build(neck)
  68. self.bbox_head = MODELS.build(bbox_head)
  69. self._init_layers()
  70. @abstractmethod
  71. def _init_layers(self) -> None:
  72. """Initialize layers except for backbone, neck and bbox_head."""
  73. pass
  74. def loss(self, batch_inputs: Tensor,
  75. batch_data_samples: SampleList) -> Union[dict, list]:
  76. """Calculate losses from a batch of inputs and data samples.
  77. Args:
  78. batch_inputs (Tensor): Input images of shape (bs, dim, H, W).
  79. These should usually be mean centered and std scaled.
  80. batch_data_samples (List[:obj:`DetDataSample`]): The batch
  81. data samples. It usually includes information such
  82. as `gt_instance` or `gt_panoptic_seg` or `gt_sem_seg`.
  83. Returns:
  84. dict: A dictionary of loss components
  85. """
  86. img_feats = self.extract_feat(batch_inputs)
  87. head_inputs_dict = self.forward_transformer(img_feats,
  88. batch_data_samples)
  89. losses = self.bbox_head.loss(
  90. **head_inputs_dict, batch_data_samples=batch_data_samples)
  91. return losses
  92. def predict(self,
  93. batch_inputs: Tensor,
  94. batch_data_samples: SampleList,
  95. rescale: bool = True) -> SampleList:
  96. """Predict results from a batch of inputs and data samples with post-
  97. processing.
  98. Args:
  99. batch_inputs (Tensor): Inputs, has shape (bs, dim, H, W).
  100. batch_data_samples (List[:obj:`DetDataSample`]): The batch
  101. data samples. It usually includes information such
  102. as `gt_instance` or `gt_panoptic_seg` or `gt_sem_seg`.
  103. rescale (bool): Whether to rescale the results.
  104. Defaults to True.
  105. Returns:
  106. list[:obj:`DetDataSample`]: Detection results of the input images.
  107. Each DetDataSample usually contain 'pred_instances'. And the
  108. `pred_instances` usually contains following keys.
  109. - scores (Tensor): Classification scores, has a shape
  110. (num_instance, )
  111. - labels (Tensor): Labels of bboxes, has a shape
  112. (num_instances, ).
  113. - bboxes (Tensor): Has a shape (num_instances, 4),
  114. the last dimension 4 arrange as (x1, y1, x2, y2).
  115. """
  116. img_feats = self.extract_feat(batch_inputs)
  117. head_inputs_dict = self.forward_transformer(img_feats,
  118. batch_data_samples)
  119. results_list = self.bbox_head.predict(
  120. **head_inputs_dict,
  121. rescale=rescale,
  122. batch_data_samples=batch_data_samples)
  123. batch_data_samples = self.add_pred_to_datasample(
  124. batch_data_samples, results_list)
  125. return batch_data_samples
  126. def _forward(
  127. self,
  128. batch_inputs: Tensor,
  129. batch_data_samples: OptSampleList = None) -> Tuple[List[Tensor]]:
  130. """Network forward process. Usually includes backbone, neck and head
  131. forward without any post-processing.
  132. Args:
  133. batch_inputs (Tensor): Inputs, has shape (bs, dim, H, W).
  134. batch_data_samples (List[:obj:`DetDataSample`], optional): The
  135. batch data samples. It usually includes information such
  136. as `gt_instance` or `gt_panoptic_seg` or `gt_sem_seg`.
  137. Defaults to None.
  138. Returns:
  139. tuple[Tensor]: A tuple of features from ``bbox_head`` forward.
  140. """
  141. img_feats = self.extract_feat(batch_inputs)
  142. head_inputs_dict = self.forward_transformer(img_feats,
  143. batch_data_samples)
  144. results = self.bbox_head.forward(**head_inputs_dict)
  145. return results
  146. def forward_transformer(self,
  147. img_feats: Tuple[Tensor],
  148. batch_data_samples: OptSampleList = None) -> Dict:
  149. """Forward process of Transformer, which includes four steps:
  150. 'pre_transformer' -> 'encoder' -> 'pre_decoder' -> 'decoder'. We
  151. summarized the parameters flow of the existing DETR-like detector,
  152. which can be illustrated as follow:
  153. .. code:: text
  154. img_feats & batch_data_samples
  155. |
  156. V
  157. +-----------------+
  158. | pre_transformer |
  159. +-----------------+
  160. | |
  161. | V
  162. | +-----------------+
  163. | | forward_encoder |
  164. | +-----------------+
  165. | |
  166. | V
  167. | +---------------+
  168. | | pre_decoder |
  169. | +---------------+
  170. | | |
  171. V V |
  172. +-----------------+ |
  173. | forward_decoder | |
  174. +-----------------+ |
  175. | |
  176. V V
  177. head_inputs_dict
  178. Args:
  179. img_feats (tuple[Tensor]): Tuple of feature maps from neck. Each
  180. feature map has shape (bs, dim, H, W).
  181. batch_data_samples (list[:obj:`DetDataSample`], optional): The
  182. batch data samples. It usually includes information such
  183. as `gt_instance` or `gt_panoptic_seg` or `gt_sem_seg`.
  184. Defaults to None.
  185. Returns:
  186. dict: The dictionary of bbox_head function inputs, which always
  187. includes the `hidden_states` of the decoder output and may contain
  188. `references` including the initial and intermediate references.
  189. """
  190. encoder_inputs_dict, decoder_inputs_dict = self.pre_transformer(
  191. img_feats, batch_data_samples)
  192. encoder_outputs_dict = self.forward_encoder(**encoder_inputs_dict)
  193. tmp_dec_in, head_inputs_dict = self.pre_decoder(**encoder_outputs_dict)
  194. decoder_inputs_dict.update(tmp_dec_in)
  195. decoder_outputs_dict = self.forward_decoder(**decoder_inputs_dict)
  196. head_inputs_dict.update(decoder_outputs_dict)
  197. return head_inputs_dict
  198. def extract_feat(self, batch_inputs: Tensor) -> Tuple[Tensor]:
  199. """Extract features.
  200. Args:
  201. batch_inputs (Tensor): Image tensor, has shape (bs, dim, H, W).
  202. Returns:
  203. tuple[Tensor]: Tuple of feature maps from neck. Each feature map
  204. has shape (bs, dim, H, W).
  205. """
  206. x = self.backbone(batch_inputs)
  207. if self.with_neck:
  208. x = self.neck(x)
  209. return x
  210. @abstractmethod
  211. def pre_transformer(
  212. self,
  213. img_feats: Tuple[Tensor],
  214. batch_data_samples: OptSampleList = None) -> Tuple[Dict, Dict]:
  215. """Process image features before feeding them to the transformer.
  216. Args:
  217. img_feats (tuple[Tensor]): Tuple of feature maps from neck. Each
  218. feature map has shape (bs, dim, H, W).
  219. batch_data_samples (list[:obj:`DetDataSample`], optional): The
  220. batch data samples. It usually includes information such
  221. as `gt_instance` or `gt_panoptic_seg` or `gt_sem_seg`.
  222. Defaults to None.
  223. Returns:
  224. tuple[dict, dict]: The first dict contains the inputs of encoder
  225. and the second dict contains the inputs of decoder.
  226. - encoder_inputs_dict (dict): The keyword args dictionary of
  227. `self.forward_encoder()`, which includes 'feat', 'feat_mask',
  228. 'feat_pos', and other algorithm-specific arguments.
  229. - decoder_inputs_dict (dict): The keyword args dictionary of
  230. `self.forward_decoder()`, which includes 'memory_mask', and
  231. other algorithm-specific arguments.
  232. """
  233. pass
  234. @abstractmethod
  235. def forward_encoder(self, feat: Tensor, feat_mask: Tensor,
  236. feat_pos: Tensor, **kwargs) -> Dict:
  237. """Forward with Transformer encoder.
  238. Args:
  239. feat (Tensor): Sequential features, has shape (bs, num_feat_points,
  240. dim).
  241. feat_mask (Tensor): ByteTensor, the padding mask of the features,
  242. has shape (bs, num_feat_points).
  243. feat_pos (Tensor): The positional embeddings of the features, has
  244. shape (bs, num_feat_points, dim).
  245. Returns:
  246. dict: The dictionary of encoder outputs, which includes the
  247. `memory` of the encoder output and other algorithm-specific
  248. arguments.
  249. """
  250. pass
  251. @abstractmethod
  252. def pre_decoder(self, memory: Tensor, **kwargs) -> Tuple[Dict, Dict]:
  253. """Prepare intermediate variables before entering Transformer decoder,
  254. such as `query`, `query_pos`, and `reference_points`.
  255. Args:
  256. memory (Tensor): The output embeddings of the Transformer encoder,
  257. has shape (bs, num_feat_points, dim).
  258. Returns:
  259. tuple[dict, dict]: The first dict contains the inputs of decoder
  260. and the second dict contains the inputs of the bbox_head function.
  261. - decoder_inputs_dict (dict): The keyword dictionary args of
  262. `self.forward_decoder()`, which includes 'query', 'query_pos',
  263. 'memory', and other algorithm-specific arguments.
  264. - head_inputs_dict (dict): The keyword dictionary args of the
  265. bbox_head functions, which is usually empty, or includes
  266. `enc_outputs_class` and `enc_outputs_class` when the detector
  267. support 'two stage' or 'query selection' strategies.
  268. """
  269. pass
  270. @abstractmethod
  271. def forward_decoder(self, query: Tensor, query_pos: Tensor, memory: Tensor,
  272. **kwargs) -> Dict:
  273. """Forward with Transformer decoder.
  274. Args:
  275. query (Tensor): The queries of decoder inputs, has shape
  276. (bs, num_queries, dim).
  277. query_pos (Tensor): The positional queries of decoder inputs,
  278. has shape (bs, num_queries, dim).
  279. memory (Tensor): The output embeddings of the Transformer encoder,
  280. has shape (bs, num_feat_points, dim).
  281. Returns:
  282. dict: The dictionary of decoder outputs, which includes the
  283. `hidden_states` of the decoder output, `references` including
  284. the initial and intermediate reference_points, and other
  285. algorithm-specific arguments.
  286. """
  287. pass