nvlm_d.py 6.39 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
5
6
7
8
9
# adapted from https://huggingface.co/nvidia/NVLM-D-72B/blob/main/modeling_nvlm_d.py
# --------------------------------------------------------
# NVLM-D
# Copyright (c) 2024 NVIDIA
# Licensed under Apache 2.0 License [see LICENSE for details]
# --------------------------------------------------------
10
from collections.abc import Mapping, Sequence
11

12
import torch
13
14
15
import torch.nn as nn
from transformers import PretrainedConfig

16
from vllm.config.multimodal import BaseDummyOptions
17
from vllm.model_executor.layers.quantization import QuantizationConfig
18
from vllm.multimodal import MULTIMODAL_REGISTRY
19
from vllm.multimodal.inputs import MultiModalDataDict, MultiModalKwargsItems
20
21
22
23
24
25
26
27
28
29
from vllm.multimodal.parse import (
    ImageEmbeddingItems,
    ImageProcessorItems,
    MultiModalDataItems,
)
from vllm.multimodal.processing import (
    PromptReplacement,
    PromptUpdate,
    PromptUpdateDetails,
)
30
from vllm.transformers_utils.processors.nvlm_d import IMG_PAD, NVLMProcessor
31
32

from .intern_vit import InternVisionModel
33
34
35
36
37
38
from .internvl import (
    BaseInternVLDummyInputsBuilder,
    BaseInternVLMultiModalProcessor,
    BaseInternVLProcessingInfo,
    InternVLChatModel,
)
39

40
41

class NVLMProcessingInfo(BaseInternVLProcessingInfo):
42
    def get_hf_processor(self, **kwargs: object) -> NVLMProcessor:
43
44
45
46
47
        return self.ctx.init_processor(
            NVLMProcessor,
            config=self.get_hf_config(),
            tokenizer=self.get_tokenizer(),
            **kwargs,
48
49
50
        )


51
class NVLMDummyInputsBuilder(BaseInternVLDummyInputsBuilder[NVLMProcessingInfo]):
52
53
54
55
56
57
58
59
    def get_dummy_text(self, mm_counts: Mapping[str, int]) -> str:
        num_images = mm_counts.get("image", 0)

        # The newline is necessary to separate ">" of the current item
        # and "<" of the next item
        return "<image>\n" * num_images

    def get_dummy_mm_data(
60
61
62
        self,
        seq_len: int,
        mm_counts: Mapping[str, int],
63
        mm_options: Mapping[str, BaseDummyOptions],
64
    ) -> MultiModalDataDict:
65
        target_width, target_height = self.info.get_image_size_with_most_features()
66
67
        num_images = mm_counts.get("image", 0)

68
        image_overrides = mm_options.get("image")
69

70
        return {
71
72
73
74
75
76
            "image": self._get_dummy_images(
                width=target_width,
                height=target_height,
                num_images=num_images,
                overrides=image_overrides,
            )
77
78
        }

79

80
class NVLMMultiModalProcessor(BaseInternVLMultiModalProcessor[NVLMProcessingInfo]):
81
    def _get_prompt_updates(
82
83
84
        self,
        mm_items: MultiModalDataItems,
        hf_processor_mm_kwargs: Mapping[str, object],
85
        out_mm_kwargs: MultiModalKwargsItems,
86
    ) -> Sequence[PromptUpdate]:
87
88
        hf_processor = self.info.get_hf_processor(**hf_processor_mm_kwargs)

89
90
91
        out_mm_data = out_mm_kwargs.get_data()
        if "image_num_patches" in out_mm_data:
            image_num_patches = out_mm_data["image_num_patches"]
92
93
            assert isinstance(image_num_patches, torch.Tensor)
            image_num_patches = image_num_patches.tolist()
94
        elif "image_embeds" in out_mm_data:
95
96
            # TODO: Use image size information in dictionary embedding inputs
            # to compute num_patches (similar to Qwen2-VL)
97
            image_num_patches = [None] * len(out_mm_data["image_embeds"])
98
99
100
101
102
        else:
            image_num_patches = []

        def get_replacement_nvlm(item_idx: int):
            images = mm_items.get_items(
103
104
                "image", (ImageEmbeddingItems, ImageProcessorItems)
            )
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119

            if isinstance(images, ImageEmbeddingItems):
                feature_size = images.get_feature_size(item_idx)
            else:
                image_size = images.get_image_size(item_idx)
                feature_size = self.info.get_num_image_tokens(
                    image_width=image_size.width,
                    image_height=image_size.height,
                    processor=hf_processor,
                )

            num_patches = image_num_patches[item_idx]
            if num_patches is not None:
                assert isinstance(num_patches, int)

120
121
            repl = hf_processor.get_image_repl(feature_size, num_patches)

122
            return PromptUpdateDetails.select_text(repl.full + "\n", IMG_PAD)
123
124
125
126
127
128
129
130
131
132
133

        # See note in dummy data regarding why we have the extra newline
        return [
            PromptReplacement(
                modality="image",
                target="<image>\n",
                replacement=get_replacement_nvlm,
            )
        ]


134
135
136
137
138
@MULTIMODAL_REGISTRY.register_processor(
    NVLMMultiModalProcessor,
    info=NVLMProcessingInfo,
    dummy_inputs=NVLMDummyInputsBuilder,
)
139
class NVLM_D_Model(InternVLChatModel):
140
    def _init_mlp1(self, config: PretrainedConfig) -> nn.Module:
141
142
143
144
145
        vit_hidden_size = config.vision_config.hidden_size
        llm_intermediate_size = config.text_config.intermediate_size
        llm_hidden_size = config.text_config.hidden_size

        return nn.Sequential(
146
147
148
149
150
151
            nn.LayerNorm(vit_hidden_size * int(1 / self.downsample_ratio) ** 2),
            nn.Linear(
                vit_hidden_size * int(1 / self.downsample_ratio) ** 2,
                llm_intermediate_size,
                bias=False,
            ),
152
153
154
155
            nn.GELU(),
            nn.Linear(llm_intermediate_size, llm_hidden_size, bias=False),
        )

156
157
158
    def _init_vision_model(
        self,
        config: PretrainedConfig,
159
        quant_config: QuantizationConfig | None,
160
161
162
163
164
165
166
        *,
        is_mono: bool,
        prefix: str,
    ):
        if not is_mono:
            vision_feature_layer = config.select_layer
            if vision_feature_layer < 0:
167
168
169
                num_hidden_layers = (
                    config.vision_config.num_hidden_layers + vision_feature_layer + 1
                )
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
            else:
                num_hidden_layers = vision_feature_layer + 1

            # We added additional dummy heads to the original num of heads to
            # make the number of heads divisible by 8.
            return InternVisionModel(
                config.vision_config,
                quant_config=quant_config,
                num_hidden_layers_override=num_hidden_layers,
                num_dummy_heads=7,
                prefix=prefix,
            )
        else:
            msg = "Monolith mode is not applicable to NVLM_D"
            raise NotImplementedError(msg)