Skip to content

Commit b681eb8

Browse files
authored
update evaluation logic (#90)
update evaluation logic
1 parent a141bd3 commit b681eb8

File tree

18 files changed

+569
-74
lines changed

18 files changed

+569
-74
lines changed

.github/ISSUE_TEMPLATE/bug_report.md

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -20,7 +20,7 @@ Steps to reproduce the behavior:
2020
**Operating environment(运行环境):**
2121
- python version [e.g. 3.6, 3.7, 3.8]
2222
- tensorflow version [e.g. 1.9.0, 1.14.0, 2.5.0]
23-
- deepmatch version [e.g. 0.3.0,]
23+
- deepmatch version [e.g. 0.3.1,]
2424

2525
**Additional context**
2626
Add any other context about the problem here.

.github/ISSUE_TEMPLATE/question.md

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -17,4 +17,4 @@ Add any other context about the problem here.
1717
**Operating environment(运行环境):**
1818
- python version [e.g. 3.6, 3.7, 3.8]
1919
- tensorflow version [e.g. 1.9.0, 1.14.0, 2.5.0]
20-
- deepmatch version [e.g. 0.3.0,]
20+
- deepmatch version [e.g. 0.3.1,]

.github/workflows/ci.yml

Lines changed: 19 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -17,7 +17,7 @@ jobs:
1717
timeout-minutes: 120
1818
strategy:
1919
matrix:
20-
python-version: [3.6,3.7,3.8]
20+
python-version: [3.6,3.7,3.8,3.9,3.10.7]
2121
tf-version: [1.9.0,1.14.0,2.5.0]
2222

2323
exclude:
@@ -57,12 +57,28 @@ jobs:
5757
tf-version: 2.8.0
5858
- python-version: 3.6
5959
tf-version: 2.9.0
60+
- python-version: 3.6
61+
tf-version: 2.10.0
6062
- python-version: 3.9
6163
tf-version: 1.4.0
64+
- python-version: 3.9
65+
tf-version: 1.9.0
6266
- python-version: 3.9
6367
tf-version: 1.15.0
6468
- python-version: 3.9
65-
tf-version: 2.2.0
69+
tf-version: 1.14.0
70+
- python-version: 3.10.7
71+
tf-version: 1.4.0
72+
- python-version: 3.10.7
73+
tf-version: 1.9.0
74+
- python-version: 3.10.7
75+
tf-version: 1.15.0
76+
- python-version: 3.10.7
77+
tf-version: 1.14.0
78+
- python-version: 3.10.7
79+
tf-version: 2.5.0
80+
- python-version: 3.10.7
81+
tf-version: 2.6.0
6682

6783
steps:
6884

@@ -75,6 +91,7 @@ jobs:
7591

7692
- name: Install dependencies
7793
run: |
94+
sudo apt update && sudo apt install -y pkg-config libhdf5-dev
7895
pip3 install -q tensorflow==${{ matrix.tf-version }}
7996
pip install -q protobuf==3.19.0
8097
pip install -q requests

.gitignore

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1,3 +1,4 @@
1+
.idea
12
# Byte-compiled / optimized / DLL files
23
__pycache__/
34
*.py[cod]

deepmatch/__init__.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,4 @@
11
from .utils import check_version
22

3-
__version__ = '0.3.0'
3+
__version__ = '0.3.1'
44
check_version(__version__)

deepmatch/models/comirec.py

Lines changed: 18 additions & 20 deletions
Original file line numberDiff line numberDiff line change
@@ -13,18 +13,20 @@
1313
from deepctr.layers.utils import NoMask, combined_dnn_input, add_func
1414
from tensorflow.python.keras.layers import Concatenate, Lambda
1515
from tensorflow.python.keras.models import Model
16+
1617
from ..inputs import create_embedding_matrix
17-
from ..layers.core import CapsuleLayer, PoolingLayer, LabelAwareAttention, SampledSoftmaxLayer, EmbeddingIndex
18+
from ..layers.core import CapsuleLayer, PoolingLayer, MaskUserEmbedding, LabelAwareAttention, SampledSoftmaxLayer, \
19+
EmbeddingIndex
1820
from ..layers.interaction import SoftmaxWeightedSum
1921
from ..utils import get_item_embedding
2022

2123

22-
def tile_user_otherfeat(user_other_feature, interest_num):
23-
return tf.tile(tf.expand_dims(user_other_feature, -2), [1, interest_num, 1])
24+
def tile_user_otherfeat(user_other_feature, k_max):
25+
return tf.tile(tf.expand_dims(user_other_feature, -2), [1, k_max, 1])
2426

2527

26-
def tile_user_his_mask(hist_len, seq_max_len, interest_num):
27-
return tf.tile(tf.sequence_mask(hist_len, seq_max_len), [1, interest_num, 1])
28+
def tile_user_his_mask(hist_len, seq_max_len, k_max):
29+
return tf.tile(tf.sequence_mask(hist_len, seq_max_len), [1, k_max, 1])
2830

2931

3032
def softmax_Weighted_Sum(input):
@@ -37,20 +39,19 @@ def softmax_Weighted_Sum(input):
3739
return high_capsule
3840

3941

40-
def ComiRec(user_feature_columns, item_feature_columns, interest_num=2, p=100, interest_extractor='sa', add_pos=False,
42+
def ComiRec(user_feature_columns, item_feature_columns, k_max=2, p=100, interest_extractor='sa',
43+
add_pos=True,
4144
user_dnn_hidden_units=(64, 32), dnn_activation='relu', dnn_use_bn=False, l2_reg_dnn=0,
4245
l2_reg_embedding=1e-6,
4346
dnn_dropout=0, output_activation='linear', sampler_config=None, seed=1024):
4447
"""Instantiates the ComiRec Model architecture.
4548
4649
:param user_feature_columns: An iterable containing user's features used by the model.
4750
:param item_feature_columns: An iterable containing item's features used by the model.
48-
:param num_sampled: int, the number of classes to randomly sample per batch.
49-
:param interest_num: int, the max size of user interest embedding
51+
:param k_max: int, the max size of user interest embedding
5052
:param p: float,the parameter for adjusting the attention distribution in LabelAwareAttention.
5153
:param interest_extractor: string, type of a multi-interest extraction module, 'sa' means self-attentive and 'dr' means dynamic routing
5254
:param add_pos: bool. Whether use positional encoding layer
53-
:param dnn_use_bn: bool. Whether use BatchNormalization before activation or not in deep net
5455
:param user_dnn_hidden_units: list,list of positive integer or empty list, the layer number and units in each layer of user tower
5556
:param dnn_activation: Activation function to use in deep net
5657
:param dnn_use_bn: bool. Whether use BatchNormalization before activation or not in deep net
@@ -131,29 +132,25 @@ def ComiRec(user_feature_columns, item_feature_columns, interest_num=2, p=100, i
131132
if interest_extractor.lower() == 'dr':
132133
high_capsule = CapsuleLayer(input_units=item_embedding_dim,
133134
out_units=item_embedding_dim, max_len=seq_max_len,
134-
k_max=interest_num)((history_emb, hist_len))
135+
k_max=k_max)((history_emb, hist_len))
135136
elif interest_extractor.lower() == 'sa':
136137
history_emb_add_pos = history_emb
137138
if add_pos:
138139
position_embedding = PositionEncoding()(history_emb)
139140
history_emb_add_pos = add_func([history_emb_add_pos, position_embedding]) # [None, max_len, emb_dim]
140141

141-
attn = DNN((item_embedding_dim * 4, interest_num), activation='tanh', l2_reg=l2_reg_dnn,
142+
attn = DNN((item_embedding_dim * 4, k_max), activation='tanh', l2_reg=l2_reg_dnn,
142143
dropout_rate=dnn_dropout, use_bn=dnn_use_bn, output_activation=None, seed=seed,
143144
name="user_dnn_attn")(history_emb_add_pos)
144-
mask = Lambda(tile_user_his_mask, arguments={'interest_num': interest_num,
145+
mask = Lambda(tile_user_his_mask, arguments={'k_max': k_max,
145146
'seq_max_len': seq_max_len})(
146-
hist_len) # [None, interest_num, max_len]
147-
# high_capsule = SoftmaxWeightedSum(dropout_rate=0, future_binding=False,
148-
# seed=seed)([attn, history_emb_add_pos, mask])
147+
hist_len) # [None, k_max, max_len]
148+
149149
high_capsule = Lambda(softmax_Weighted_Sum)((history_emb_add_pos, mask, attn))
150150

151-
print("high_capsule",
152-
high_capsule) # Tensor("softmax_weighted_sum/MatMul:0", shape=(None, 2, 32), dtype=float32) Tensor("capsule_layer/Reshape_1:0", shape=(None, 2, 32), dtype=float32)
153151
if len(dnn_input_emb_list) > 0 or len(dense_value_list) > 0:
154152
user_other_feature = combined_dnn_input(dnn_input_emb_list, dense_value_list)
155-
other_feature_tile = Lambda(tile_user_otherfeat, arguments={'interest_num': interest_num})(user_other_feature)
156-
print("other_feature_tile", other_feature_tile, "NoMask", NoMask()(other_feature_tile))
153+
other_feature_tile = Lambda(tile_user_otherfeat, arguments={'k_max': k_max})(user_other_feature)
157154
user_deep_input = Concatenate()([NoMask()(other_feature_tile), high_capsule])
158155
else:
159156
user_deep_input = high_capsule
@@ -173,7 +170,8 @@ def ComiRec(user_feature_columns, item_feature_columns, interest_num=2, p=100, i
173170

174171
pooling_item_embedding_weight = PoolingLayer()([item_embedding_weight])
175172

176-
user_embedding_final = LabelAwareAttention(k_max=interest_num, pow_p=p)((user_embeddings, target_emb))
173+
user_embedding_final = LabelAwareAttention(k_max=k_max, pow_p=p)((user_embeddings, target_emb))
174+
177175
output = SampledSoftmaxLayer(sampler_config._asdict())(
178176
[pooling_item_embedding_weight, user_embedding_final, item_features[item_feature_name]])
179177
model = Model(inputs=inputs_list + item_inputs_list, outputs=output)

deepmatch/models/mind.py

Lines changed: 1 addition & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -53,11 +53,9 @@ def MIND(user_feature_columns, item_feature_columns, k_max=2, p=100, dynamic_k=F
5353
5454
:param user_feature_columns: An iterable containing user's features used by the model.
5555
:param item_feature_columns: An iterable containing item's features used by the model.
56-
:param num_sampled: int, the number of classes to randomly sample per batch.
5756
:param k_max: int, the max size of user interest embedding
5857
:param p: float,the parameter for adjusting the attention distribution in LabelAwareAttention.
5958
:param dynamic_k: bool, whether or not use dynamic interest number
60-
:param dnn_use_bn: bool. Whether use BatchNormalization before activation or not in deep net
6159
:param user_dnn_hidden_units: list,list of positive integer or empty list, the layer number and units in each layer of user tower
6260
:param dnn_activation: Activation function to use in deep net
6361
:param dnn_use_bn: bool. Whether use BatchNormalization before activation or not in deep net
@@ -169,7 +167,7 @@ def MIND(user_feature_columns, item_feature_columns, k_max=2, p=100, dynamic_k=F
169167
user_embedding_final = LabelAwareAttention(k_max=k_max, pow_p=p)((user_embeddings, target_emb, interest_num))
170168
else:
171169
user_embedding_final = LabelAwareAttention(k_max=k_max, pow_p=p)((user_embeddings, target_emb))
172-
print("swc")
170+
173171
output = SampledSoftmaxLayer(sampler_config._asdict())(
174172
[pooling_item_embedding_weight, user_embedding_final, item_features[item_feature_name]])
175173
model = Model(inputs=inputs_list + item_inputs_list, outputs=output)

deepmatch/models/sdm.py

Lines changed: 0 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -30,7 +30,6 @@ def SDM(user_feature_columns, item_feature_columns, history_feature_list, units=
3030
:param user_feature_columns: An iterable containing user's features used by the model.
3131
:param item_feature_columns: An iterable containing item's features used by the model.
3232
:param history_feature_list: list,to indicate short and prefer sequence sparse field
33-
:param num_sampled: int, the number of classes to randomly sample per batch.
3433
:param units: int, dimension for each output layer
3534
:param rnn_layers: int, layer number of rnn
3635
:param dropout_rate: float in [0,1), the probability we will drop out a given DNN coordinate.

deepmatch/models/youtubednn.py

Lines changed: 0 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -23,7 +23,6 @@ def YoutubeDNN(user_feature_columns, item_feature_columns,
2323
2424
:param user_feature_columns: An iterable containing user's features used by the model.
2525
:param item_feature_columns: An iterable containing item's features used by the model.
26-
:param num_sampled: int, the number of classes to randomly sample per batch.
2726
:param user_dnn_hidden_units: list,list of positive integer or empty list, the layer number and units in each layer of user tower
2827
:param dnn_activation: Activation function to use in deep net
2928
:param dnn_use_bn: bool. Whether use BatchNormalization before activation or not in deep net

docs/source/History.md

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,5 @@
11
# History
2+
- 10/31/2022 : [v0.3.1](https://github.com/shenweichen/DeepMatch/releases/tag/v0.3.1) released.Add `ComiRec` model .
23
- 07/04/2022 : [v0.3.0](https://github.com/shenweichen/DeepMatch/releases/tag/v0.3.0) released.Support different negative sampling strategies, including `inbatch`, `uniform`, `frequency`, `adaptive`.
34
- 06/17/2022 : [v0.2.1](https://github.com/shenweichen/DeepMatch/releases/tag/v0.2.1) released.Fix some bugs.
45
- 10/12/2020 : [v0.2.0](https://github.com/shenweichen/DeepMatch/releases/tag/v0.2.0) released.Support different initializers for different embedding weights and loading pretrained embeddings.

0 commit comments

Comments
 (0)