Skip to content

Commit 451d0cb

Browse files
committed
Created SoftmaxCrossEntropyLoss source files for ONNX FE
1 parent 9380826 commit 451d0cb

File tree

2 files changed

+65
-65
lines changed

2 files changed

+65
-65
lines changed

src/frontends/onnx/frontend/src/op/softmax_cross_entropy_loss.hpp

Lines changed: 5 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -10,11 +10,11 @@ namespace ov {
1010
namespace frontend {
1111
namespace onnx {
1212
namespace ai_onnx{
13-
namespace opset_12 {
14-
OutputVector softmax_cross_entropy_loss(const Node& node);
15-
} // namespace opset_12
16-
namespace opset_13 {
17-
OutputVector softmax_cross_entropy_loss(const Node& node);
13+
namespace opset_12 {
14+
OutputVector softmax_cross_entropy_loss(const Node& node);
15+
} // namespace opset_12
16+
namespace opset_13 {
17+
OutputVector softmax_cross_entropy_loss(const Node& node);
1818
} // namespace opset_13
1919
} // namespace ai_onnx
2020
} // namespace onnx

src/frontends/onnx/frontend/src/op/softmax_crossentropy_loss.cpp

Lines changed: 60 additions & 60 deletions
Original file line numberDiff line numberDiff line change
@@ -3,93 +3,93 @@
33
//
44

55
#include "core/operator_set.hpp"
6+
#include "openvino/op/divide.hpp"
67
#include "openvino/op/gather.hpp"
78
#include "openvino/op/log.hpp"
9+
#include "openvino/op/multiply.hpp"
810
#include "openvino/op/negative.hpp"
911
#include "openvino/op/reduce_mean.hpp"
1012
#include "openvino/op/reduce_sum.hpp"
1113
#include "openvino/op/softmax.hpp"
1214
#include "softmax_cross_entropy_loss.hpp"
13-
#include "openvino/op/multiply.hpp"
14-
#include "openvino/op/divide.hpp"
1515

1616
namespace ov {
1717
namespace frontend {
1818
namespace onnx {
1919
namespace {
20-
OutputVector impl_softmax_cross_entropy(const Node& node, int64_t axis_default) {
21-
const auto inputs = node.get_ov_inputs();
20+
OutputVector impl_softmax_cross_entropy(const Node& node, int64_t axis_default) {
21+
const auto inputs = node.get_ov_inputs();
2222

23-
const auto scores = inputs[0];
24-
const auto labels = inputs[1];
23+
const auto scores = inputs[0];
24+
const auto labels = inputs[1];
2525

26-
// Optional weights paramater
27-
bool has_weights = inputs.size() > 2;
28-
std::shared_ptr<ov::Node> weights_gather = nullptr;
26+
// Optional weights paramater
27+
bool has_weights = inputs.size() > 2;
28+
std::shared_ptr<ov::Node> weights_gather = nullptr;
2929

30-
if (has_weights) {
31-
const auto weights = inputs[2];
32-
const auto axis_for_weights = ov::op::v0::Constant::create(element::i64, {}, {0});
33-
weights_gather = std::make_shared<ov::op::v8::Gather>(weights, labels, axis_for_weights);
34-
}
30+
if (has_weights) {
31+
const auto weights = inputs[2];
32+
const auto axis_for_weights = ov::op::v0::Constant::create(element::i64, {}, {0});
33+
weights_gather = std::make_shared<ov::op::v8::Gather>(weights, labels, axis_for_weights);
34+
}
3535

36-
// Getting attributes for axis and reduction mode
37-
const auto axis = node.get_attribute_value<int64_t>("axis", axis_default);
38-
const auto reduction = node.get_attribute_value<std::string>("reduction", "mean");
36+
// Getting attributes for axis and reduction mode
37+
const auto axis = node.get_attribute_value<int64_t>("axis", axis_default);
38+
const auto reduction = node.get_attribute_value<std::string>("reduction", "mean");
3939

40-
// Computing softmax & it's logarithm
41-
const auto softmax = std::make_shared<ov::op::v8::Softmax>(scores, axis);
42-
const auto log_softmax = std::make_shared<ov::op::v0::Log>(softmax);
40+
// Computing softmax & it's logarithm
41+
const auto softmax = std::make_shared<ov::op::v8::Softmax>(scores, axis);
42+
const auto log_softmax = std::make_shared<ov::op::v0::Log>(softmax);
4343

44-
const auto axis_const = ov::op::v0::Constant::create(element::i64, {}, {axis});
45-
const auto gathered = std::make_shared<ov::op::v8::Gather>(log_softmax, labels, axis_const);
44+
const auto axis_const = ov::op::v0::Constant::create(element::i64, {}, {axis});
45+
const auto gathered = std::make_shared<ov::op::v8::Gather>(log_softmax, labels, axis_const);
4646

47+
// Computing loss
48+
std::shared_ptr<ov::Node> loss = std::make_shared<ov::op::v0::Negative>(gathered);
4749

48-
// Computing loss
49-
std::shared_ptr<ov::Node> loss = std::make_shared<ov::op::v0::Negative>(gathered);
50+
if (has_weights) {
51+
loss = std::make_shared<ov::op::v1::Multiply>(loss, weights_gather);
52+
}
5053

51-
if (has_weights) {
52-
loss = std::make_shared<ov::op::v1::Multiply>(loss, weights_gather);
53-
}
54+
// applying reduction as mentioned in
55+
// https://github.com/onnx/onnx/blob/main/docs/Changelog.md#softmaxcrossentropyloss-12
5456

55-
// applying reduction as mentioned in https://github.com/onnx/onnx/blob/main/docs/Changelog.md#softmaxcrossentropyloss-12
56-
57-
if (reduction != "None") {
58-
// Reduce over the axis corresponding to each sample
59-
// Reducing over axis 0, assuming the loss tensor shape is [batch_size]
60-
const auto reduce_axis = ov::op::v0::Constant::create(ov::element::i64, {1}, {0});
57+
if (reduction != "None") {
58+
// Reduce over the axis corresponding to each sample
59+
// Reducing over axis 0, assuming the loss tensor shape is [batch_size]
60+
const auto reduce_axis = ov::op::v0::Constant::create(ov::element::i64, {1}, {0});
6161

62-
if (reduction == "mean") {
63-
if (has_weights) {
64-
auto loss_sum = std::make_shared<ov::op::v1::ReduceSum>(loss->output(0), reduce_axis, true);
65-
auto weight_sum = std::make_shared<ov::op::v1::ReduceSum>(weights_gather->output(0), reduce_axis, true);
66-
loss = std::make_shared<ov::op::v1::Divide>(loss_sum, weight_sum);
67-
} else {
68-
loss = std::make_shared<ov::op::v1::ReduceMean>(loss->output(0), reduce_axis, true);
69-
}
70-
} else if (reduction == "sum") {
71-
loss = std::make_shared<ov::op::v1::ReduceSum>(loss->output(0), reduce_axis, true);
62+
if (reduction == "mean") {
63+
if (has_weights) {
64+
auto loss_sum = std::make_shared<ov::op::v1::ReduceSum>(loss->output(0), reduce_axis, true);
65+
auto weight_sum = std::make_shared<ov::op::v1::ReduceSum>(weights_gather->output(0), reduce_axis, true);
66+
loss = std::make_shared<ov::op::v1::Divide>(loss_sum, weight_sum);
67+
} else {
68+
loss = std::make_shared<ov::op::v1::ReduceMean>(loss->output(0), reduce_axis, true);
7269
}
70+
} else if (reduction == "sum") {
71+
loss = std::make_shared<ov::op::v1::ReduceSum>(loss->output(0), reduce_axis, true);
7372
}
74-
75-
return {loss};
7673
}
74+
75+
return {loss};
7776
}
77+
} // namespace
7878
namespace ai_onnx {
79-
namespace opset_12 {
80-
OutputVector ov::frontend::onnx::ai_onnx::opset_12::softmax_cross_entropy_loss(const Node& node) {
81-
return impl_softmax_cross_entropy(node, 1);
82-
}
83-
ONNX_OP("SoftmaxCrossEntropyLoss", OPSET_SINCE(12), ai_onnx::opset_12::softmax_cross_entropy_loss);
84-
}
85-
namespace opset_13 {
86-
OutputVector ov::frontend::onnx::ai_onnx::opset_13::softmax_cross_entropy_loss(const Node& node) {
87-
return impl_softmax_cross_entropy(node, 1);
88-
}
89-
90-
ONNX_OP("SoftmaxCrossEntropyLoss", OPSET_SINCE(13), ai_onnx::opset_13::softmax_cross_entropy_loss);
91-
}
92-
}
79+
namespace opset_12 {
80+
OutputVector ov::frontend::onnx::ai_onnx::opset_12::softmax_cross_entropy_loss(const Node& node) {
81+
return impl_softmax_cross_entropy(node, 1);
9382
}
83+
ONNX_OP("SoftmaxCrossEntropyLoss", OPSET_SINCE(12), ai_onnx::opset_12::softmax_cross_entropy_loss);
84+
} // namespace opset_12
85+
namespace opset_13 {
86+
OutputVector ov::frontend::onnx::ai_onnx::opset_13::softmax_cross_entropy_loss(const Node& node) {
87+
return impl_softmax_cross_entropy(node, 1);
9488
}
95-
}
89+
90+
ONNX_OP("SoftmaxCrossEntropyLoss", OPSET_SINCE(13), ai_onnx::opset_13::softmax_cross_entropy_loss);
91+
} // namespace opset_13
92+
} // namespace ai_onnx
93+
} // namespace onnx
94+
} // namespace frontend
95+
} // namespace ov

0 commit comments

Comments
 (0)