dino_layers.py 26 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552
  1. # Copyright (c) OpenMMLab. All rights reserved.
  2. import warnings
  3. from typing import Tuple, Union
  4. import torch
  5. from mmengine.model import BaseModule
  6. from torch import Tensor, nn
  7. from mmdet.structures import SampleList
  8. from mmdet.structures.bbox import bbox_xyxy_to_cxcywh
  9. from mmdet.utils import OptConfigType
  10. from .deformable_detr_layers import DeformableDetrTransformerDecoder
  11. from .utils import MLP, coordinate_to_encoding, inverse_sigmoid
  12. class DinoTransformerDecoder(DeformableDetrTransformerDecoder):
  13. """Transformer encoder of DINO."""
  14. def _init_layers(self) -> None:
  15. """Initialize decoder layers."""
  16. super()._init_layers()
  17. self.ref_point_head = MLP(self.embed_dims * 2, self.embed_dims,
  18. self.embed_dims, 2)
  19. self.norm = nn.LayerNorm(self.embed_dims)
  20. def forward(self, query: Tensor, value: Tensor, key_padding_mask: Tensor,
  21. self_attn_mask: Tensor, reference_points: Tensor,
  22. spatial_shapes: Tensor, level_start_index: Tensor,
  23. valid_ratios: Tensor, reg_branches: nn.ModuleList,
  24. **kwargs) -> Tensor:
  25. """Forward function of Transformer encoder.
  26. Args:
  27. query (Tensor): The input query, has shape (num_queries, bs, dim).
  28. value (Tensor): The input values, has shape (num_value, bs, dim).
  29. key_padding_mask (Tensor): The `key_padding_mask` of `self_attn`
  30. input. ByteTensor, has shape (num_queries, bs).
  31. self_attn_mask (Tensor): The attention mask to prevent information
  32. leakage from different denoising groups and matching parts, has
  33. shape (num_queries_total, num_queries_total). It is `None` when
  34. `self.training` is `False`.
  35. reference_points (Tensor): The initial reference, has shape
  36. (bs, num_queries, 4) with the last dimension arranged as
  37. (cx, cy, w, h).
  38. spatial_shapes (Tensor): Spatial shapes of features in all levels,
  39. has shape (num_levels, 2), last dimension represents (h, w).
  40. level_start_index (Tensor): The start index of each level.
  41. A tensor has shape (num_levels, ) and can be represented
  42. as [0, h_0*w_0, h_0*w_0+h_1*w_1, ...].
  43. valid_ratios (Tensor): The ratios of the valid width and the valid
  44. height relative to the width and the height of features in all
  45. levels, has shape (bs, num_levels, 2).
  46. reg_branches: (obj:`nn.ModuleList`): Used for refining the
  47. regression results.
  48. Returns:
  49. Tensor: Output queries of Transformer encoder, which is also
  50. called 'encoder output embeddings' or 'memory', has shape
  51. (num_queries, bs, dim)
  52. """
  53. intermediate = []
  54. intermediate_reference_points = [reference_points]
  55. for lid, layer in enumerate(self.layers):
  56. if reference_points.shape[-1] == 4:
  57. reference_points_input = \
  58. reference_points[:, :, None] * torch.cat(
  59. [valid_ratios, valid_ratios], -1)[:, None]
  60. else:
  61. assert reference_points.shape[-1] == 2
  62. reference_points_input = \
  63. reference_points[:, :, None] * valid_ratios[:, None]
  64. query_sine_embed = coordinate_to_encoding(
  65. reference_points_input[:, :, 0, :])
  66. query_pos = self.ref_point_head(query_sine_embed)
  67. query = layer(
  68. query,
  69. query_pos=query_pos,
  70. value=value,
  71. key_padding_mask=key_padding_mask,
  72. self_attn_mask=self_attn_mask,
  73. spatial_shapes=spatial_shapes,
  74. level_start_index=level_start_index,
  75. valid_ratios=valid_ratios,
  76. reference_points=reference_points_input,
  77. **kwargs)
  78. if reg_branches is not None:
  79. tmp = reg_branches[lid](query)
  80. assert reference_points.shape[-1] == 4
  81. new_reference_points = tmp + inverse_sigmoid(
  82. reference_points, eps=1e-3)
  83. new_reference_points = new_reference_points.sigmoid()
  84. reference_points = new_reference_points.detach()
  85. if self.return_intermediate:
  86. intermediate.append(self.norm(query))
  87. intermediate_reference_points.append(new_reference_points)
  88. # NOTE this is for the "Look Forward Twice" module,
  89. # in the DeformDETR, reference_points was appended.
  90. if self.return_intermediate:
  91. return torch.stack(intermediate), torch.stack(
  92. intermediate_reference_points)
  93. return query, reference_points
  94. class CdnQueryGenerator(BaseModule):
  95. """Implement query generator of the Contrastive denoising (CDN) proposed in
  96. `DINO: DETR with Improved DeNoising Anchor Boxes for End-to-End Object
  97. Detection <https://arxiv.org/abs/2203.03605>`_
  98. Code is modified from the `official github repo
  99. <https://github.com/IDEA-Research/DINO>`_.
  100. Args:
  101. num_classes (int): Number of object classes.
  102. embed_dims (int): The embedding dimensions of the generated queries.
  103. num_matching_queries (int): The queries number of the matching part.
  104. Used for generating dn_mask.
  105. label_noise_scale (float): The scale of label noise, defaults to 0.5.
  106. box_noise_scale (float): The scale of box noise, defaults to 1.0.
  107. group_cfg (:obj:`ConfigDict` or dict, optional): The config of the
  108. denoising queries grouping, includes `dynamic`, `num_dn_queries`,
  109. and `num_groups`. Two grouping strategies, 'static dn groups' and
  110. 'dynamic dn groups', are supported. When `dynamic` is `False`,
  111. the `num_groups` should be set, and the number of denoising query
  112. groups will always be `num_groups`. When `dynamic` is `True`, the
  113. `num_dn_queries` should be set, and the group number will be
  114. dynamic to ensure that the denoising queries number will not exceed
  115. `num_dn_queries` to prevent large fluctuations of memory. Defaults
  116. to `None`.
  117. """
  118. def __init__(self,
  119. num_classes: int,
  120. embed_dims: int,
  121. num_matching_queries: int,
  122. label_noise_scale: float = 0.5,
  123. box_noise_scale: float = 1.0,
  124. group_cfg: OptConfigType = None) -> None:
  125. super().__init__()
  126. self.num_classes = num_classes
  127. self.embed_dims = embed_dims
  128. self.num_matching_queries = num_matching_queries
  129. self.label_noise_scale = label_noise_scale
  130. self.box_noise_scale = box_noise_scale
  131. # prepare grouping strategy
  132. group_cfg = {} if group_cfg is None else group_cfg
  133. self.dynamic_dn_groups = group_cfg.get('dynamic', True)
  134. if self.dynamic_dn_groups:
  135. if 'num_dn_queries' not in group_cfg:
  136. warnings.warn("'num_dn_queries' should be set when using "
  137. 'dynamic dn groups, use 100 as default.')
  138. self.num_dn_queries = group_cfg.get('num_dn_queries', 100)
  139. assert isinstance(self.num_dn_queries, int), \
  140. f'Expected the num_dn_queries to have type int, but got ' \
  141. f'{self.num_dn_queries}({type(self.num_dn_queries)}). '
  142. else:
  143. assert 'num_groups' in group_cfg, \
  144. 'num_groups should be set when using static dn groups'
  145. self.num_groups = group_cfg['num_groups']
  146. assert isinstance(self.num_groups, int), \
  147. f'Expected the num_groups to have type int, but got ' \
  148. f'{self.num_groups}({type(self.num_groups)}). '
  149. # NOTE The original repo of DINO set the num_embeddings 92 for coco,
  150. # 91 (0~90) of which represents target classes and the 92 (91)
  151. # indicates `Unknown` class. However, the embedding of `unknown` class
  152. # is not used in the original DINO.
  153. # TODO: num_classes + 1 or num_classes ?
  154. self.label_embedding = nn.Embedding(self.num_classes, self.embed_dims)
  155. def __call__(self, batch_data_samples: SampleList) -> tuple:
  156. """Generate contrastive denoising (cdn) queries with ground truth.
  157. Descriptions of the Number Values in code and comments:
  158. - num_target_total: the total target number of the input batch
  159. samples.
  160. - max_num_target: the max target number of the input batch samples.
  161. - num_noisy_targets: the total targets number after adding noise,
  162. i.e., num_target_total * num_groups * 2.
  163. - num_denoising_queries: the length of the output batched queries,
  164. i.e., max_num_target * num_groups * 2.
  165. NOTE The format of input bboxes in batch_data_samples is unnormalized
  166. (x, y, x, y), and the output bbox queries are embedded by normalized
  167. (cx, cy, w, h) format bboxes going through inverse_sigmoid.
  168. Args:
  169. batch_data_samples (list[:obj:`DetDataSample`]): List of the batch
  170. data samples, each includes `gt_instance` which has attributes
  171. `bboxes` and `labels`. The `bboxes` has unnormalized coordinate
  172. format (x, y, x, y).
  173. Returns:
  174. tuple: The outputs of the dn query generator.
  175. - dn_label_query (Tensor): The output content queries for denoising
  176. part, has shape (bs, num_denoising_queries, dim), where
  177. `num_denoising_queries = max_num_target * num_groups * 2`.
  178. - dn_bbox_query (Tensor): The output reference bboxes as positions
  179. of queries for denoising part, which are embedded by normalized
  180. (cx, cy, w, h) format bboxes going through inverse_sigmoid, has
  181. shape (bs, num_denoising_queries, 4) with the last dimension
  182. arranged as (cx, cy, w, h).
  183. - attn_mask (Tensor): The attention mask to prevent information
  184. leakage from different denoising groups and matching parts,
  185. will be used as `self_attn_mask` of the `decoder`, has shape
  186. (num_queries_total, num_queries_total), where `num_queries_total`
  187. is the sum of `num_denoising_queries` and `num_matching_queries`.
  188. - dn_meta (Dict[str, int]): The dictionary saves information about
  189. group collation, including 'num_denoising_queries' and
  190. 'num_denoising_groups'. It will be used for split outputs of
  191. denoising and matching parts and loss calculation.
  192. """
  193. # normalize bbox and collate ground truth (gt)
  194. gt_labels_list = []
  195. gt_bboxes_list = []
  196. for sample in batch_data_samples:
  197. img_h, img_w = sample.img_shape
  198. bboxes = sample.gt_instances.bboxes
  199. factor = bboxes.new_tensor([img_w, img_h, img_w,
  200. img_h]).unsqueeze(0)
  201. bboxes_normalized = bboxes / factor
  202. gt_bboxes_list.append(bboxes_normalized)
  203. gt_labels_list.append(sample.gt_instances.labels)
  204. gt_labels = torch.cat(gt_labels_list) # (num_target_total, 4)
  205. gt_bboxes = torch.cat(gt_bboxes_list)
  206. num_target_list = [len(bboxes) for bboxes in gt_bboxes_list]
  207. max_num_target = max(num_target_list)
  208. num_groups = self.get_num_groups(max_num_target)
  209. dn_label_query = self.generate_dn_label_query(gt_labels, num_groups)
  210. dn_bbox_query = self.generate_dn_bbox_query(gt_bboxes, num_groups)
  211. # The `batch_idx` saves the batch index of the corresponding sample
  212. # for each target, has shape (num_target_total).
  213. batch_idx = torch.cat([
  214. torch.full_like(t.long(), i) for i, t in enumerate(gt_labels_list)
  215. ])
  216. dn_label_query, dn_bbox_query = self.collate_dn_queries(
  217. dn_label_query, dn_bbox_query, batch_idx, len(batch_data_samples),
  218. num_groups)
  219. attn_mask = self.generate_dn_mask(
  220. max_num_target, num_groups, device=dn_label_query.device)
  221. dn_meta = dict(
  222. num_denoising_queries=int(max_num_target * 2 * num_groups),
  223. num_denoising_groups=num_groups)
  224. return dn_label_query, dn_bbox_query, attn_mask, dn_meta
  225. def get_num_groups(self, max_num_target: int = None) -> int:
  226. """Calculate denoising query groups number.
  227. Two grouping strategies, 'static dn groups' and 'dynamic dn groups',
  228. are supported. When `self.dynamic_dn_groups` is `False`, the number
  229. of denoising query groups will always be `self.num_groups`. When
  230. `self.dynamic_dn_groups` is `True`, the group number will be dynamic,
  231. ensuring the denoising queries number will not exceed
  232. `self.num_dn_queries` to prevent large fluctuations of memory.
  233. NOTE The `num_group` is shared for different samples in a batch. When
  234. the target numbers in the samples varies, the denoising queries of the
  235. samples containing fewer targets are padded to the max length.
  236. Args:
  237. max_num_target (int, optional): The max target number of the batch
  238. samples. It will only be used when `self.dynamic_dn_groups` is
  239. `True`. Defaults to `None`.
  240. Returns:
  241. int: The denoising group number of the current batch.
  242. """
  243. if self.dynamic_dn_groups:
  244. assert max_num_target is not None, \
  245. 'group_queries should be provided when using ' \
  246. 'dynamic dn groups'
  247. if max_num_target == 0:
  248. num_groups = 1
  249. else:
  250. num_groups = self.num_dn_queries // max_num_target
  251. else:
  252. num_groups = self.num_groups
  253. if num_groups < 1:
  254. num_groups = 1
  255. return int(num_groups)
  256. def generate_dn_label_query(self, gt_labels: Tensor,
  257. num_groups: int) -> Tensor:
  258. """Generate noisy labels and their query embeddings.
  259. The strategy for generating noisy labels is: Randomly choose labels of
  260. `self.label_noise_scale * 0.5` proportion and override each of them
  261. with a random object category label.
  262. NOTE Not add noise to all labels. Besides, the `self.label_noise_scale
  263. * 0.5` arg is the ratio of the chosen positions, which is higher than
  264. the actual proportion of noisy labels, because the labels to override
  265. may be correct. And the gap becomes larger as the number of target
  266. categories decreases. The users should notice this and modify the scale
  267. arg or the corresponding logic according to specific dataset.
  268. Args:
  269. gt_labels (Tensor): The concatenated gt labels of all samples
  270. in the batch, has shape (num_target_total, ) where
  271. `num_target_total = sum(num_target_list)`.
  272. num_groups (int): The number of denoising query groups.
  273. Returns:
  274. Tensor: The query embeddings of noisy labels, has shape
  275. (num_noisy_targets, embed_dims), where `num_noisy_targets =
  276. num_target_total * num_groups * 2`.
  277. """
  278. assert self.label_noise_scale > 0
  279. gt_labels_expand = gt_labels.repeat(2 * num_groups,
  280. 1).view(-1) # Note `* 2` # noqa
  281. p = torch.rand_like(gt_labels_expand.float())
  282. chosen_indice = torch.nonzero(p < (self.label_noise_scale * 0.5)).view(
  283. -1) # Note `* 0.5`
  284. new_labels = torch.randint_like(chosen_indice, 0, self.num_classes)
  285. noisy_labels_expand = gt_labels_expand.scatter(0, chosen_indice,
  286. new_labels)
  287. dn_label_query = self.label_embedding(noisy_labels_expand)
  288. return dn_label_query
  289. def generate_dn_bbox_query(self, gt_bboxes: Tensor,
  290. num_groups: int) -> Tensor:
  291. """Generate noisy bboxes and their query embeddings.
  292. The strategy for generating noisy bboxes is as follow:
  293. .. code:: text
  294. +--------------------+
  295. | negative |
  296. | +----------+ |
  297. | | positive | |
  298. | | +-----|----+------------+
  299. | | | | | |
  300. | +----+-----+ | |
  301. | | | |
  302. +---------+----------+ |
  303. | |
  304. | gt bbox |
  305. | |
  306. | +---------+----------+
  307. | | | |
  308. | | +----+-----+ |
  309. | | | | | |
  310. +-------------|--- +----+ | |
  311. | | positive | |
  312. | +----------+ |
  313. | negative |
  314. +--------------------+
  315. The random noise is added to the top-left and down-right point
  316. positions, hence, normalized (x, y, x, y) format of bboxes are
  317. required. The noisy bboxes of positive queries have the points
  318. both within the inner square, while those of negative queries
  319. have the points both between the inner and outer squares.
  320. Besides, the length of outer square is twice as long as that of
  321. the inner square, i.e., self.box_noise_scale * w_or_h / 2.
  322. NOTE The noise is added to all the bboxes. Moreover, there is still
  323. unconsidered case when one point is within the positive square and
  324. the others is between the inner and outer squares.
  325. Args:
  326. gt_bboxes (Tensor): The concatenated gt bboxes of all samples
  327. in the batch, has shape (num_target_total, 4) with the last
  328. dimension arranged as (cx, cy, w, h) where
  329. `num_target_total = sum(num_target_list)`.
  330. num_groups (int): The number of denoising query groups.
  331. Returns:
  332. Tensor: The output noisy bboxes, which are embedded by normalized
  333. (cx, cy, w, h) format bboxes going through inverse_sigmoid, has
  334. shape (num_noisy_targets, 4) with the last dimension arranged as
  335. (cx, cy, w, h), where
  336. `num_noisy_targets = num_target_total * num_groups * 2`.
  337. """
  338. assert self.box_noise_scale > 0
  339. device = gt_bboxes.device
  340. # expand gt_bboxes as groups
  341. gt_bboxes_expand = gt_bboxes.repeat(2 * num_groups, 1) # xyxy
  342. # obtain index of negative queries in gt_bboxes_expand
  343. positive_idx = torch.arange(
  344. len(gt_bboxes), dtype=torch.long, device=device)
  345. positive_idx = positive_idx.unsqueeze(0).repeat(num_groups, 1)
  346. positive_idx += 2 * len(gt_bboxes) * torch.arange(
  347. num_groups, dtype=torch.long, device=device)[:, None]
  348. positive_idx = positive_idx.flatten()
  349. negative_idx = positive_idx + len(gt_bboxes)
  350. # determine the sign of each element in the random part of the added
  351. # noise to be positive or negative randomly.
  352. rand_sign = torch.randint_like(
  353. gt_bboxes_expand, low=0, high=2,
  354. dtype=torch.float32) * 2.0 - 1.0 # [low, high), 1 or -1, randomly
  355. # calculate the random part of the added noise
  356. rand_part = torch.rand_like(gt_bboxes_expand) # [0, 1)
  357. rand_part[negative_idx] += 1.0 # pos: [0, 1); neg: [1, 2)
  358. rand_part *= rand_sign # pos: (-1, 1); neg: (-2, -1] U [1, 2)
  359. # add noise to the bboxes
  360. bboxes_whwh = bbox_xyxy_to_cxcywh(gt_bboxes_expand)[:, 2:].repeat(1, 2)
  361. noisy_bboxes_expand = gt_bboxes_expand + torch.mul(
  362. rand_part, bboxes_whwh) * self.box_noise_scale / 2 # xyxy
  363. noisy_bboxes_expand = noisy_bboxes_expand.clamp(min=0.0, max=1.0)
  364. noisy_bboxes_expand = bbox_xyxy_to_cxcywh(noisy_bboxes_expand)
  365. dn_bbox_query = inverse_sigmoid(noisy_bboxes_expand, eps=1e-3)
  366. return dn_bbox_query
  367. def collate_dn_queries(self, input_label_query: Tensor,
  368. input_bbox_query: Tensor, batch_idx: Tensor,
  369. batch_size: int, num_groups: int) -> Tuple[Tensor]:
  370. """Collate generated queries to obtain batched dn queries.
  371. The strategy for query collation is as follow:
  372. .. code:: text
  373. input_queries (num_target_total, query_dim)
  374. P_A1 P_B1 P_B2 N_A1 N_B1 N_B2 P'A1 P'B1 P'B2 N'A1 N'B1 N'B2
  375. |________ group1 ________| |________ group2 ________|
  376. |
  377. V
  378. P_A1 Pad0 N_A1 Pad0 P'A1 Pad0 N'A1 Pad0
  379. P_B1 P_B2 N_B1 N_B2 P'B1 P'B2 N'B1 N'B2
  380. |____ group1 ____| |____ group2 ____|
  381. batched_queries (batch_size, max_num_target, query_dim)
  382. where query_dim is 4 for bbox and self.embed_dims for label.
  383. Notation: _-group 1; '-group 2;
  384. A-Sample1(has 1 target); B-sample2(has 2 targets)
  385. Args:
  386. input_label_query (Tensor): The generated label queries of all
  387. targets, has shape (num_target_total, embed_dims) where
  388. `num_target_total = sum(num_target_list)`.
  389. input_bbox_query (Tensor): The generated bbox queries of all
  390. targets, has shape (num_target_total, 4) with the last
  391. dimension arranged as (cx, cy, w, h).
  392. batch_idx (Tensor): The batch index of the corresponding sample
  393. for each target, has shape (num_target_total).
  394. batch_size (int): The size of the input batch.
  395. num_groups (int): The number of denoising query groups.
  396. Returns:
  397. tuple[Tensor]: Output batched label and bbox queries.
  398. - batched_label_query (Tensor): The output batched label queries,
  399. has shape (batch_size, max_num_target, embed_dims).
  400. - batched_bbox_query (Tensor): The output batched bbox queries,
  401. has shape (batch_size, max_num_target, 4) with the last dimension
  402. arranged as (cx, cy, w, h).
  403. """
  404. device = input_label_query.device
  405. num_target_list = [
  406. torch.sum(batch_idx == idx) for idx in range(batch_size)
  407. ]
  408. max_num_target = max(num_target_list)
  409. num_denoising_queries = int(max_num_target * 2 * num_groups)
  410. map_query_index = torch.cat([
  411. torch.arange(num_target, device=device)
  412. for num_target in num_target_list
  413. ])
  414. map_query_index = torch.cat([
  415. map_query_index + max_num_target * i for i in range(2 * num_groups)
  416. ]).long()
  417. batch_idx_expand = batch_idx.repeat(2 * num_groups, 1).view(-1)
  418. mapper = (batch_idx_expand, map_query_index)
  419. batched_label_query = torch.zeros(
  420. batch_size, num_denoising_queries, self.embed_dims, device=device)
  421. batched_bbox_query = torch.zeros(
  422. batch_size, num_denoising_queries, 4, device=device)
  423. batched_label_query[mapper] = input_label_query
  424. batched_bbox_query[mapper] = input_bbox_query
  425. return batched_label_query, batched_bbox_query
  426. def generate_dn_mask(self, max_num_target: int, num_groups: int,
  427. device: Union[torch.device, str]) -> Tensor:
  428. """Generate attention mask to prevent information leakage from
  429. different denoising groups and matching parts.
  430. .. code:: text
  431. 0 0 0 0 1 1 1 1 0 0 0 0 0
  432. 0 0 0 0 1 1 1 1 0 0 0 0 0
  433. 0 0 0 0 1 1 1 1 0 0 0 0 0
  434. 0 0 0 0 1 1 1 1 0 0 0 0 0
  435. 1 1 1 1 0 0 0 0 0 0 0 0 0
  436. 1 1 1 1 0 0 0 0 0 0 0 0 0
  437. 1 1 1 1 0 0 0 0 0 0 0 0 0
  438. 1 1 1 1 0 0 0 0 0 0 0 0 0
  439. 1 1 1 1 1 1 1 1 0 0 0 0 0
  440. 1 1 1 1 1 1 1 1 0 0 0 0 0
  441. 1 1 1 1 1 1 1 1 0 0 0 0 0
  442. 1 1 1 1 1 1 1 1 0 0 0 0 0
  443. 1 1 1 1 1 1 1 1 0 0 0 0 0
  444. max_num_target |_| |_________| num_matching_queries
  445. |_____________| num_denoising_queries
  446. 1 -> True (Masked), means 'can not see'.
  447. 0 -> False (UnMasked), means 'can see'.
  448. Args:
  449. max_num_target (int): The max target number of the input batch
  450. samples.
  451. num_groups (int): The number of denoising query groups.
  452. device (obj:`device` or str): The device of generated mask.
  453. Returns:
  454. Tensor: The attention mask to prevent information leakage from
  455. different denoising groups and matching parts, will be used as
  456. `self_attn_mask` of the `decoder`, has shape (num_queries_total,
  457. num_queries_total), where `num_queries_total` is the sum of
  458. `num_denoising_queries` and `num_matching_queries`.
  459. """
  460. num_denoising_queries = int(max_num_target * 2 * num_groups)
  461. num_queries_total = num_denoising_queries + self.num_matching_queries
  462. attn_mask = torch.zeros(
  463. num_queries_total,
  464. num_queries_total,
  465. device=device,
  466. dtype=torch.bool)
  467. # Make the matching part cannot see the denoising groups
  468. attn_mask[num_denoising_queries:, :num_denoising_queries] = True
  469. # Make the denoising groups cannot see each other
  470. for i in range(num_groups):
  471. # Mask rows of one group per step.
  472. row_scope = slice(max_num_target * 2 * i,
  473. max_num_target * 2 * (i + 1))
  474. left_scope = slice(max_num_target * 2 * i)
  475. right_scope = slice(max_num_target * 2 * (i + 1),
  476. num_denoising_queries)
  477. attn_mask[row_scope, right_scope] = True
  478. attn_mask[row_scope, left_scope] = True
  479. return attn_mask