Skip to content

Conversation

@sbucaille
Copy link
Contributor

What does this PR do?

Implements RF-DETR

Fixes #36879

Before submitting

  • This PR fixes a typo or improves the docs (you can dismiss the other checks if that's the case).
  • Did you read the contributor guideline,
    Pull Request section?
  • Was this discussed/approved via a Github issue or the forum? Please add a link
    to it if that's the case.
  • Did you make sure to update the documentation with your changes? Here are the
    documentation guidelines, and
    here are tips on formatting docstrings.
  • Did you write any new necessary tests?

Who can review?

@qubvel

@sbucaille
Copy link
Contributor Author

Just a small message to present the architecture and what it looks like from 🤗 transformers point of view :
image

RF-DETR is based on LW-DETR and DeformableDETR. The LW-DETR is based on DETR but modified the encoder to be a ViT instead of a CNN (like ResNet) and they added the appropriate MultiScaleProjector to make the link between the encoder and the decoder. RF-DETR changed in LW-DETR the encoder from a ViT to DinoV2WithRegisters with a "window" mechanism as well as changed the classical DETR decoder by a DeformableDETR decoder.

There is basically 2 things to write :

  • The RFDetrMultiScaleProjector which was originally implemented in LWDetr which RFDetr is based on but not present in the library.
  • The RFDetrBackbone with the underlying classes built on top of DinoV2WithRegisters

One difficulty I may see in advance is the following :
Am I right saying that there should be only one XXPreTrainedModel class per modeling file ?
In our case, we will need to create a RFDetrBackbone class which requires a PreTrainedModel to be considered as an AutoBackbone, presumably inheritating from DinoV2WithRegistersPreTrainedModel. We will also need to create RFDetrModel and RFDetrForObjectDetection which both require a PreTrainedModel, which will likely inherit from DeformableDetrPreTrainedModel.
If yes then I need to "merge" both classes but _supports_flash_attn_2 is not the same for both, DinoV2 supports it but not DeformableDetr.

I noticed your PR about refactoring attention in ViTs, is there any plan for other models such as Detr, RTDetr etc to add FlashAttention ?
I guess for now I'll just set _supports_flash_attn_2 as false.

Let me know what you guys think

@qubvel
Copy link
Contributor

qubvel commented Mar 22, 2025

Hi @sbucaille, thanks for the detailed write-up!

Am I right saying that there should be only one XXPreTrainedModel class per modeling file ? In our case, we will need to create a RFDetrBackbone class which requires a PreTrainedModel to be considered as an AutoBackbone, presumably inheritating from DinoV2WithRegistersPreTrainedModel. We will also need to create RFDetrModel and RFDetrForObjectDetection which both require a PreTrainedModel, which will likely inherit from DeformableDetrPreTrainedModel.

We can add DinoV2WithRegistersBackbone class directly into the dino_v2_with_registers model, would it work?

I noticed your #36545 about refactoring attention in ViTs, is there any plan for other models such as Detr, RTDetr etc to add FlashAttention ?

Not at the moment, from my experiments it was not required for detr-based models and did not give any speedup. However, it might be more relevant for transformer-based encoder. Let's keep it simple initially and set it to False as you suggested

@sbucaille
Copy link
Contributor Author

We can't use the DinoV2WithRegistersBackbone as there is the "window" mechanism in the middle of the forward methods, here is an example (not final) :

class RFDetrBackboneLayer(Dinov2WithRegistersLayer):
    def __init__(self, config):
        super(Dinov2WithRegistersLayer).__init__(config)

        self.num_windows = config.num_windows

    def forward(
        self,
        hidden_states: torch.Tensor,
        head_mask: Optional[torch.Tensor] = None,
        output_attentions: bool = False,
        run_full_attention: bool = False,
    ) -> Union[Tuple[torch.Tensor, torch.Tensor], Tuple[torch.Tensor]]:
        assert head_mask is None, "head_mask is not supported for windowed attention"
        assert not output_attentions, "output_attentions is not supported for windowed attention"
        shortcut = hidden_states
        if run_full_attention:
            # reshape x to remove windows
            B, HW, C = hidden_states.shape
            num_windows_squared = self.num_windows**2
            hidden_states = hidden_states.view(B // num_windows_squared, num_windows_squared * HW, C)

        self_attention_outputs = self.attention(
            self.norm1(hidden_states),  # in Dinov2WithRegisters, layernorm is applied before self-attention
            head_mask,
            output_attentions=output_attentions,
        )
        attention_output = self_attention_outputs[0]

        if run_full_attention:
            # reshape x to add windows back
            B, HW, C = hidden_states.shape
            num_windows_squared = self.num_windows**2
            # hidden_states = hidden_states.view(B * num_windows_squared, HW // num_windows_squared, C)
            attention_output = attention_output.view(B * num_windows_squared, HW // num_windows_squared, C)

        attention_output = self.layer_scale1(attention_output)
        outputs = self_attention_outputs[1:]  # add self attentions if we output attention weights

        # first residual connection
        hidden_states = self.drop_path(attention_output) + shortcut

        # in Dinov2WithRegisters, layernorm is also applied after self-attention
        layer_output = self.norm2(hidden_states)
        layer_output = self.mlp(layer_output)
        layer_output = self.layer_scale2(layer_output)

        # second residual connection
        layer_output = self.drop_path(layer_output) + hidden_states

        outputs = (layer_output,) + outputs

        return outputs

That's why I think we necessarily need a custom Backbone class for that 🤔

@qubvel
Copy link
Contributor

qubvel commented Mar 24, 2025

Hmm, am I correct that this part was added?

if run_full_attention:
            # reshape x to add windows back
            B, HW, C = hidden_states.shape
            num_windows_squared = self.num_windows**2
            # hidden_states = hidden_states.view(B * num_windows_squared, HW // num_windows_squared, C)
            attention_output = attention_output.view(B * num_windows_squared, HW // num_windows_squared, C)

It looks like it is a reshape only operation, we can return attention_output as is, and reshape all layers output later, right?

@sbucaille
Copy link
Contributor Author

You are right, but it is not the only example. I'll stick to my original plan until I have something running with actual results and I'll take care of refactoring this part later, I'll ping you when it's ready.

@sbucaille
Copy link
Contributor Author

Hey @qubvel, in the end I made modeling files follow the rt_detr folder structure with modeling_rf_detr_dinov2_with_registers.py being like the modeling_rt_detr_resnet.py where the backbone is defined and modeling_rf_detr.py like modeling_rt_detr.py where the encoder/decoder is defined with the use of any possible backbone.
What do you think ? I'll continue tonight

Also I had issues with the modular mechanism where utils/modular_model_converter.py always enforced the name RfDetr... instead of RFDetr... as I wanted to follow the naming convention used for RTDetr. So I ended up using the copied from mechanism.

@qubvel
Copy link
Contributor

qubvel commented Mar 25, 2025

Hey, let's use RfDert name + modular, it's ok! RfDetr is a correct naming format while RTDetr is an exception made before modular was introduced

@sbucaille
Copy link
Contributor Author

Ok sorry I confused the problems I had, I didn't have a problem with the capital letters of Rf or RF but rather a problem with the prefix RfDetr used in modular file with DeformableDetr
Using :

class RfDetrModel(DeformableDetrModel):
    pass

generates a bunch of Rf... classes like RfConvEncoder instead of RfDetrConvEncoder. I supposed RfDetr and DeformableDetr share Detr in their name which the modular script failing, I don't have the case when using RfDetrDinov2WithRegisters(Dinov2WithRegisters) naturally. The way I can avoid this is by forcing the naming of these classes by overwriting the __init__ method like this :

class RfDetrConvEncoder(DeformableDetrConvEncoder):
    pass

class RfDetrModel(DeformableDetrModel):
    def __init__(self, config: RfDetrConfig):
        super().__init__(config)

        backbone = RfDetrConvEncoder(config)
        ...

But the problem also appears for ModelOutput's, so I'm forced to rewrite the whole forward methods for many classes, which makes using modular a bit useless in my opinion...
So for modeling_rf_detr_dinov2_with_registers.py, I can keep a modular file but not for modeling_rf_detr.py, I think I'll need to use the copied from mechanism.

Should I open an issue ? Maybe @ArthurZucker have some insights on this problem ?

@qubvel
Copy link
Contributor

qubvel commented Mar 25, 2025

cc @Cyrilvallez re modular you faced somthing similar

@konstantinos-p
Copy link
Contributor

konstantinos-p commented Apr 8, 2025

I'm also facing a similar issue with @sbucaille while working on DinoDetr.

@Cyrilvallez
Copy link
Member

Hey! Super super sorry, I missed the ping! Indeed, model sharing the last part of their names can introduce issues. I think I have an idea that should fix it, while keeping general prefix renaming sound (which is very hard in practice)! I'll try to tackle it asap and will come back to you!

@Cyrilvallez
Copy link
Member

Cyrilvallez commented Apr 28, 2025

Hey @sbucaille @konstantinos-p! It will be solved by #37829 🤗 I will merge asap! Sorry for the wait on this!

EDIT: Just merged the PR!

@ZaraCook ZaraCook mentioned this pull request Jun 18, 2025
@sbucaille sbucaille mentioned this pull request Sep 19, 2025
@sbucaille sbucaille force-pushed the add_rf_detr branch 2 times, most recently from 2b88599 to 50829ac Compare September 26, 2025 22:54
@sbucaille sbucaille marked this pull request as ready for review September 26, 2025 22:57
@sbucaille
Copy link
Contributor Author

@yonigozlan Ready for a first review, this branch is based on the LWDetr until it will be merged

@sbucaille sbucaille force-pushed the add_rf_detr branch 2 times, most recently from 46bd06b to 105f3da Compare October 22, 2025 02:23
Copy link
Member

@yonigozlan yonigozlan left a comment

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Hey @sbucaille , very nice PR! I mentioned some small things to change, but once lw detr is merged we should be able to merge this quickly!

Comment on lines 116 to 120
self.register_tokens = (
nn.Parameter(torch.zeros(1, config.num_register_tokens, config.hidden_size))
if config.num_register_tokens > 0
else None
)
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

It looks like num_register_tokens is 0 for all models in the convert file. So do we really need all the logic associated to it? Can we inherit from dinov2 instead?

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Indeed, fixed in b89a4af

Comment on lines 103 to 106
window_block_indexes = set(range(self._out_indices[-1] + 1))
window_block_indexes.difference_update(self._out_indices)
window_block_indexes = list(window_block_indexes)
self.window_block_indexes = window_block_indexes
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

this is a bit verbose and hard to read. Instead, let's hardcode the window_block_indices in the config like we do for vitdet for example

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Hmmm, in my opinion, I think it is still better this way as window_block_indices consists of the inverse set of out_indices, this way we don't have to check whether the provided window_block_indices are valid or not and raise an error if it is not the case.

batch_size * num_windows**2, num_h_patches_per_window * num_w_patches_per_window, -1
)
windowed_cls_token_with_pos_embed = cls_token_with_pos_embed.repeat(num_windows**2, 1, 1)
embeddings = torch.cat((windowed_cls_token_with_pos_embed, windowed_pixel_tokens), dim=1)
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Let's put all that in a window_partition utility, like we do for vitdet

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Done in 68c8918

def forward(
self,
hidden_states: torch.Tensor,
remove_windows: bool = False,
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

let's add a use_global_attention attribute to the layers when we instantiate them instead of passing an arg to the forward

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Refactored in 6a345f8

Comment on lines 186 to 200
if remove_windows:
# reshape x to remove windows
B, HW, C = hidden_states.shape
num_windows_squared = self.num_windows**2
hidden_states = hidden_states.view(B // num_windows_squared, num_windows_squared * HW, C)

hidden_states_norm = self.norm1(hidden_states)
self_attention_output = self.attention(hidden_states_norm)

if remove_windows:
# reshape x to add windows back
B, HW, C = hidden_states.shape
num_windows_squared = self.num_windows**2
# hidden_states = hidden_states.view(B * num_windows_squared, HW // num_windows_squared, C)
self_attention_output = self_attention_output.view(B * num_windows_squared, HW // num_windows_squared, C)
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Let's use window_partition and window_unpartition utilities here as well

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Done in 68c8918

Comment on lines 1 to 18
import io

import requests
from PIL import Image

from transformers import AutoImageProcessor, RFDetrBackbone, RFDetrConfig


images = ["https://media.roboflow.com/notebooks/examples/dog-2.jpeg"]

images = [Image.open(io.BytesIO(requests.get(url).content)) for url in images]

processor = AutoImageProcessor.from_pretrained("facebook/detr-resnet-50")
inputs = processor(images, return_tensors="pt")

config = RFDetrConfig()
backbone = RFDetrBackbone(config=config.backbone_config)
# model = RFDetrForObjectDetection.from_config()
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

To remove

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Ouupps, removed in 2151d0d

@@ -0,0 +1,357 @@
from typing import Optional
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Very nice use of modular 🤗

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Thanks, you guys have done an insane work on this feature, quality of life when adding models compared to before is on another level

def num_key_value_heads(self) -> int:
return self.decoder_self_attention_heads

@property
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Why is this needed?

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Leftovers from LW Detr, I'll remove them on the other PR so when I'll rebase in the future it will be gone

logger = logging.get_logger(__name__)


class RfDetrConfig(PretrainedConfig):
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

We can still inherit form LwDetrConfig for the properties

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Fixed in 0697a29

("llava_next_video", ("LlavaNextImageProcessor", "LlavaNextImageProcessorFast")),
("llava_next_video", ("LlavaNextVideoImageProcessor", None)),
("llava_onevision", ("LlavaOnevisionImageProcessor", "LlavaOnevisionImageProcessorFast")),
("lw_detr", ("LwDetrImageProcessor", "LwDetrImageProcessorFast")),
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

we still need a mapping for RFDETR as I mentioned in the LW-DETR PR

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Added in a6d5e2c

@sbucaille
Copy link
Contributor Author

sbucaille commented Jan 23, 2026

@molbap @vasqu nvm it's ready for a review
cc @stevhliu for the docs 🙂

I have a question regarding the WeightTransforms, should they be only in the convert script and saved on the hub without reverse mapping like I did, or should they be in the conversion_mapping file with the weights saved as the original ?

Copy link
Member

@stevhliu stevhliu left a comment

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

docs lgtm, thanks!

@github-actions
Copy link
Contributor

[For maintainers] Suggested jobs to run (before merge)

run-slow: auto, lw_detr, rf_detr

@github-actions
Copy link
Contributor

View the CircleCI Test Summary for this PR:

https://huggingface.co/spaces/transformers-community/circle-ci-viz?pr=36895&sha=1599f8

Sign up for free to join this conversation on GitHub. Already have an account? Sign in to comment

Projects

None yet

Development

Successfully merging this pull request may close these issues.

Add RF-DETR model

7 participants