Skip to content

Commit 112a9c2

Browse files
authored
Merge branch 'main' into pipeline
2 parents 3760e1e + 9bb4334 commit 112a9c2

24 files changed

+881
-124
lines changed

optimum/exporters/openvino/__main__.py

+17-4
Original file line numberDiff line numberDiff line change
@@ -13,6 +13,7 @@
1313
# limitations under the License.
1414

1515
import logging
16+
import warnings
1617
from pathlib import Path
1718
from typing import TYPE_CHECKING, Any, Callable, Dict, Optional, Union
1819

@@ -57,6 +58,7 @@ def main_export(
5758
force_download: bool = False,
5859
local_files_only: bool = False,
5960
use_auth_token: Optional[Union[bool, str]] = None,
61+
token: Optional[Union[bool, str]] = None,
6062
model_kwargs: Optional[Dict[str, Any]] = None,
6163
custom_export_configs: Optional[Dict[str, "OnnxConfig"]] = None,
6264
fn_get_submodels: Optional[Callable] = None,
@@ -107,9 +109,11 @@ def main_export(
107109
cached versions if they exist.
108110
local_files_only (`Optional[bool]`, defaults to `False`):
109111
Whether or not to only look at local files (i.e., do not try to download the model).
110-
use_auth_token (`Optional[str]`, defaults to `None`):
112+
use_auth_token (Optional[Union[bool, str]], defaults to `None`):
113+
Deprecated. Please use `token` instead.
114+
token (Optional[Union[bool, str]], defaults to `None`):
111115
The token to use as HTTP bearer authorization for remote files. If `True`, will use the token generated
112-
when running `transformers-cli login` (stored in `~/.huggingface`).
116+
when running `huggingface-cli login` (stored in `~/.huggingface`).
113117
model_kwargs (`Optional[Dict[str, Any]]`, defaults to `None`):
114118
Experimental usage: keyword arguments to pass to the model during
115119
the export. This argument should be used along the `custom_export_configs` argument
@@ -138,6 +142,15 @@ def main_export(
138142
```
139143
"""
140144

145+
if use_auth_token is not None:
146+
warnings.warn(
147+
"The `use_auth_token` argument is deprecated and will be removed soon. Please use the `token` argument instead.",
148+
FutureWarning,
149+
)
150+
if token is not None:
151+
raise ValueError("You cannot use both `use_auth_token` and `token` arguments at the same time.")
152+
token = use_auth_token
153+
141154
if compression_option is not None:
142155
logger.warning(
143156
"The `compression_option` argument is deprecated and will be removed in optimum-intel v1.17.0. "
@@ -196,7 +209,7 @@ def main_export(
196209
subfolder=subfolder,
197210
revision=revision,
198211
cache_dir=cache_dir,
199-
use_auth_token=use_auth_token,
212+
token=token,
200213
local_files_only=local_files_only,
201214
force_download=force_download,
202215
trust_remote_code=trust_remote_code,
@@ -268,7 +281,7 @@ class StoreAttr(object):
268281
subfolder=subfolder,
269282
revision=revision,
270283
cache_dir=cache_dir,
271-
use_auth_token=use_auth_token,
284+
token=token,
272285
local_files_only=local_files_only,
273286
force_download=force_download,
274287
trust_remote_code=trust_remote_code,

optimum/exporters/openvino/convert.py

+45-3
Original file line numberDiff line numberDiff line change
@@ -20,9 +20,10 @@
2020
from pathlib import Path
2121
from typing import TYPE_CHECKING, Any, Callable, Dict, List, Optional, Tuple, Union
2222

23+
import onnx
2324
from transformers.utils import is_tf_available, is_torch_available
2425

25-
from openvino.runtime import PartialShape, save_model
26+
from openvino.runtime import Model, PartialShape, save_model
2627
from openvino.runtime.exceptions import OVTypeError
2728
from openvino.runtime.utils.types import get_element_type
2829
from openvino.tools.ovc import convert_model
@@ -32,6 +33,14 @@
3233
from optimum.exporters.onnx.convert import export_pytorch as export_pytorch_to_onnx
3334
from optimum.exporters.onnx.convert import export_tensorflow as export_tensorflow_onnx
3435
from optimum.exporters.utils import _get_submodels_and_export_configs
36+
from optimum.intel.utils.import_utils import (
37+
_nncf_version,
38+
_optimum_intel_version,
39+
_optimum_version,
40+
_timm_version,
41+
_torch_version,
42+
_transformers_version,
43+
)
3544
from optimum.utils import DEFAULT_DUMMY_SHAPES, is_diffusers_available
3645
from optimum.utils.save_utils import maybe_save_preprocessors
3746

@@ -81,6 +90,8 @@ def _save_model(model, path: str, ov_config: Optional["OVConfig"] = None):
8190

8291
compress_to_fp16 = ov_config.dtype == "fp16"
8392

93+
library_name = TasksManager.infer_library_from_model(Path(path).parent)
94+
model = _add_version_info_to_model(model, library_name)
8495
save_model(model, path, compress_to_fp16)
8596

8697

@@ -347,6 +358,7 @@ def ts_patched_forward(*args, **kwargs):
347358

348359
with patcher:
349360
check_dummy_inputs_are_allowed(model, dummy_inputs)
361+
sig = inspect.signature(model.forward) if hasattr(model, "forward") else inspect.signature(model.call)
350362
inputs = config.ordered_inputs(model)
351363
input_names = list(inputs.keys())
352364
output_names = list(config.outputs.keys())
@@ -376,7 +388,6 @@ def ts_patched_forward(*args, **kwargs):
376388
ov_config=ov_config,
377389
)
378390

379-
sig = inspect.signature(model.forward) if hasattr(model, "forward") else inspect.signature(model.call)
380391
ordered_dummy_inputs = {param: dummy_inputs[param] for param in sig.parameters if param in dummy_inputs}
381392
if not ordered_dummy_inputs:
382393
ordered_dummy_inputs = dummy_inputs
@@ -392,7 +403,7 @@ def ts_patched_forward(*args, **kwargs):
392403
inp_tensor.get_tensor().set_names({input_name})
393404
inp_data = flatten_inputs[idx]
394405
static_shape = PartialShape(inp_data.shape)
395-
dims = inputs[input_name]
406+
dims = inputs.get(input_name, [])
396407
for dim in dims:
397408
static_shape[dim] = -1
398409
inp_tensor.get_node().set_partial_shape(static_shape)
@@ -689,3 +700,34 @@ def export_tokenizer(
689700

690701
for model, file_name in zip(converted, (OV_TOKENIZER_NAME, OV_DETOKENIZER_NAME)):
691702
save_model(model, output / file_name.format(suffix))
703+
704+
705+
def _add_version_info_to_model(model: Model, library_name: Optional[str] = None):
706+
"""
707+
Add dependency versions to OpenVINO model
708+
"""
709+
try:
710+
model.set_rt_info(_transformers_version, ["optimum", "transformers_version"])
711+
model.set_rt_info(_torch_version, ["optimum", "pytorch_version"])
712+
model.set_rt_info(_optimum_intel_version, ["optimum", "optimum_intel_version"])
713+
model.set_rt_info(_optimum_version, ["optimum", "optimum_version"])
714+
715+
if any("token_embeddings" in output.get_names() for output in model.outputs):
716+
import sentence_transformers
717+
718+
model.set_rt_info(sentence_transformers.__version__, ["optimum", "sentence_transformers_version"])
719+
if library_name == "diffusers":
720+
model.set_rt_info(_optimum_version, ["optimum", "diffusers_version"])
721+
elif library_name == "timm":
722+
model.set_rt_info(_timm_version, ["optimum", "timm_version"])
723+
rt_info = model.get_rt_info()
724+
if "nncf" in rt_info:
725+
model.set_rt_info(_nncf_version, ["optimum", "nncf_version"])
726+
input_model = rt_info["conversion_parameters"].get("input_model", None)
727+
if input_model is not None and "onnx" in input_model.value:
728+
model.set_rt_info(onnx.__version__, ["optimum", "onnx_version"])
729+
730+
except Exception:
731+
pass
732+
733+
return model

optimum/exporters/openvino/model_configs.py

+83-1
Original file line numberDiff line numberDiff line change
@@ -19,7 +19,16 @@
1919
from transformers.utils import is_tf_available
2020

2121
from optimum.exporters.onnx.config import TextDecoderOnnxConfig, TextDecoderWithPositionIdsOnnxConfig
22-
from optimum.exporters.onnx.model_configs import FalconOnnxConfig, GemmaOnnxConfig, LlamaOnnxConfig, PhiOnnxConfig
22+
from optimum.exporters.onnx.model_configs import (
23+
FalconOnnxConfig,
24+
GemmaOnnxConfig,
25+
LlamaOnnxConfig,
26+
MPTOnnxConfig,
27+
PhiOnnxConfig,
28+
UNetOnnxConfig,
29+
VaeDecoderOnnxConfig,
30+
VaeEncoderOnnxConfig,
31+
)
2332
from optimum.exporters.tasks import TasksManager
2433
from optimum.utils import DEFAULT_DUMMY_SHAPES
2534
from optimum.utils.input_generators import (
@@ -35,8 +44,10 @@
3544
BaichuanModelPatcher,
3645
ChatGLMModelPatcher,
3746
GemmaModelPatcher,
47+
InternLMPatcher,
3848
LlamaModelPatcher,
3949
MixtralModelPatcher,
50+
MPTModelPatcher,
4051
Phi3ModelPatcher,
4152
QwenModelPatcher,
4253
)
@@ -431,6 +442,11 @@ class InternLM2OpenVINOConfig(TextDecoderWithPositionIdsOnnxConfig):
431442
DUMMY_PKV_GENERATOR_CLASS = MistralDummyPastKeyValuesGenerator
432443
NORMALIZED_CONFIG_CLASS = NormalizedTextConfig
433444

445+
def patch_model_for_export(
446+
self, model: Union["PreTrainedModel", "TFPreTrainedModel"], model_kwargs: Optional[Dict[str, Any]] = None
447+
) -> "ModelPatcher":
448+
return InternLMPatcher(self, model, model_kwargs=model_kwargs)
449+
434450

435451
@register_in_tasks_manager("orion", *["text-generation", "text-generation-with-past"], library_name="transformers")
436452
class OrionOpenVINOConfig(TextDecoderWithPositionIdsOnnxConfig):
@@ -447,6 +463,16 @@ class OlmoOpenVINOConfig(TextDecoderWithPositionIdsOnnxConfig):
447463
NORMALIZED_CONFIG_CLASS = NormalizedTextConfig
448464

449465

466+
@register_in_tasks_manager(
467+
"mpt", *["text-generation", "text-generation-with-past", "text-classification"], library_name="transformers"
468+
)
469+
class MPTOpenVINOConfig(MPTOnnxConfig):
470+
def patch_model_for_export(
471+
self, model: Union["PreTrainedModel", "TFPreTrainedModel"], model_kwargs: Optional[Dict[str, Any]] = None
472+
) -> "ModelPatcher":
473+
return MPTModelPatcher(self, model, model_kwargs=model_kwargs)
474+
475+
450476
@register_in_tasks_manager(
451477
"phi3",
452478
*[
@@ -510,3 +536,59 @@ class FalconOpenVINOConfig(FalconOnnxConfig):
510536
OVFalconDummyPastKeyValuesGenerator,
511537
) + TextDecoderOnnxConfig.DUMMY_INPUT_GENERATOR_CLASSES
512538
DUMMY_PKV_GENERATOR_CLASS = OVFalconDummyPastKeyValuesGenerator
539+
540+
541+
@register_in_tasks_manager("unet", *["semantic-segmentation"], library_name="diffusers")
542+
class UNetOpenVINOConfig(UNetOnnxConfig):
543+
@property
544+
def inputs(self) -> Dict[str, Dict[int, str]]:
545+
common_inputs = {
546+
"sample": {0: "batch_size", 2: "height", 3: "width"},
547+
"timestep": {0: "steps"},
548+
"encoder_hidden_states": {0: "batch_size", 1: "sequence_length"},
549+
}
550+
551+
# TODO : add text_image, image and image_embeds
552+
if getattr(self._normalized_config, "addition_embed_type", None) == "text_time":
553+
common_inputs["text_embeds"] = {0: "batch_size"}
554+
common_inputs["time_ids"] = {0: "batch_size"}
555+
556+
if getattr(self._normalized_config, "time_cond_proj_dim", None) is not None:
557+
common_inputs["timestep_cond"] = {0: "batch_size"}
558+
return common_inputs
559+
560+
@property
561+
def outputs(self) -> Dict[str, Dict[int, str]]:
562+
return {
563+
"out_sample": {0: "batch_size", 2: "height", 3: "width"},
564+
}
565+
566+
567+
@register_in_tasks_manager("vae-encoder", *["semantic-segmentation"], library_name="diffusers")
568+
class VaeEncoderOpenVINOConfig(VaeEncoderOnnxConfig):
569+
@property
570+
def inputs(self) -> Dict[str, Dict[int, str]]:
571+
return {
572+
"sample": {0: "batch_size", 2: "height", 3: "width"},
573+
}
574+
575+
@property
576+
def outputs(self) -> Dict[str, Dict[int, str]]:
577+
return {
578+
"latent_sample": {0: "batch_size", 2: "height_latent", 3: "width_latent"},
579+
}
580+
581+
582+
@register_in_tasks_manager("vae-decoder", *["semantic-segmentation"], library_name="diffusers")
583+
class VaeDecoderOpenVINOConfig(VaeDecoderOnnxConfig):
584+
@property
585+
def inputs(self) -> Dict[str, Dict[int, str]]:
586+
return {
587+
"latent_sample": {0: "batch_size", 2: "height_latent", 3: "width_latent"},
588+
}
589+
590+
@property
591+
def outputs(self) -> Dict[str, Dict[int, str]]:
592+
return {
593+
"sample": {0: "batch_size", 2: "height", 3: "width"},
594+
}

0 commit comments

Comments
 (0)